Update.
[pytorch.git] / tiny_vae.py
1 #!/usr/bin/env python
2
3 # @XREMOTE_HOST: elk.fleuret.org
4 # @XREMOTE_EXEC: python
5 # @XREMOTE_PRE: source ${HOME}/misc/venv/pytorch/bin/activate
6 # @XREMOTE_PRE: ln -sf ${HOME}/data/pytorch ./data
7 # @XREMOTE_GET: *.png
8
9 # Any copyright is dedicated to the Public Domain.
10 # https://creativecommons.org/publicdomain/zero/1.0/
11
12 # Written by Francois Fleuret <francois@fleuret.org>
13
14 import sys, os, argparse, time, math, itertools
15
16 import torch, torchvision
17
18 from torch import optim, nn
19 from torch.nn import functional as F
20
21 ######################################################################
22
23 device = torch.device("cuda" if torch.cuda.is_available() else "cpu")
24
25 ######################################################################
26
27 parser = argparse.ArgumentParser(
28     description="Very simple implementation of a VAE for teaching."
29 )
30
31 parser.add_argument("--nb_epochs", type=int, default=25)
32
33 parser.add_argument("--learning_rate", type=float, default=1e-3)
34
35 parser.add_argument("--batch_size", type=int, default=100)
36
37 parser.add_argument("--data_dir", type=str, default="./data/")
38
39 parser.add_argument("--log_filename", type=str, default="train.log")
40
41 parser.add_argument("--latent_dim", type=int, default=32)
42
43 parser.add_argument("--nb_channels", type=int, default=128)
44
45 parser.add_argument("--no_dkl", action="store_true")
46
47 # With that option, do not follow the setup of the original VAE paper
48 # of forcing the variance of X|Z to 1 during training and to 0 for
49 # sampling, but optimize and use the variance.
50
51 parser.add_argument("--no_hacks", action="store_true")
52
53 args = parser.parse_args()
54
55 log_file = open(args.log_filename, "w")
56
57 ######################################################################
58
59
60 def log_string(s):
61     t = time.strftime("%Y-%m-%d_%H:%M:%S - ", time.localtime())
62
63     if log_file is not None:
64         log_file.write(t + s + "\n")
65         log_file.flush()
66
67     print(t + s)
68     sys.stdout.flush()
69
70
71 ######################################################################
72
73
74 def sample_gaussian(mu, log_var):
75     std = log_var.mul(0.5).exp()
76     return torch.randn(mu.size(), device=mu.device) * std + mu
77
78
79 def log_p_gaussian(x, mu, log_var):
80     var = log_var.exp()
81     return (
82         (-0.5 * ((x - mu).pow(2) / var) - 0.5 * log_var - 0.5 * math.log(2 * math.pi))
83         .flatten(1)
84         .sum(1)
85     )
86
87
88 def dkl_gaussians(mean_a, log_var_a, mean_b, log_var_b):
89     mean_a, log_var_a = mean_a.flatten(1), log_var_a.flatten(1)
90     mean_b, log_var_b = mean_b.flatten(1), log_var_b.flatten(1)
91     var_a = log_var_a.exp()
92     var_b = log_var_b.exp()
93     return 0.5 * (
94         log_var_b - log_var_a - 1 + (mean_a - mean_b).pow(2) / var_b + var_a / var_b
95     ).sum(1)
96
97
98 ######################################################################
99
100
101 class LatentGivenImageNet(nn.Module):
102     def __init__(self, nb_channels, latent_dim):
103         super().__init__()
104
105         self.model = nn.Sequential(
106             nn.Conv2d(1, nb_channels, kernel_size=1),  # to 28x28
107             nn.ReLU(inplace=True),
108             nn.Conv2d(nb_channels, nb_channels, kernel_size=5),  # to 24x24
109             nn.ReLU(inplace=True),
110             nn.Conv2d(nb_channels, nb_channels, kernel_size=5),  # to 20x20
111             nn.ReLU(inplace=True),
112             nn.Conv2d(nb_channels, nb_channels, kernel_size=4, stride=2),  # to 9x9
113             nn.ReLU(inplace=True),
114             nn.Conv2d(nb_channels, nb_channels, kernel_size=3, stride=2),  # to 4x4
115             nn.ReLU(inplace=True),
116             nn.Conv2d(nb_channels, 2 * latent_dim, kernel_size=4),
117         )
118
119     def forward(self, x):
120         output = self.model(x).view(x.size(0), 2, -1)
121         mu, log_var = output[:, 0], output[:, 1]
122         return mu, log_var
123
124
125 class ImageGivenLatentNet(nn.Module):
126     def __init__(self, nb_channels, latent_dim):
127         super().__init__()
128
129         self.model = nn.Sequential(
130             nn.ConvTranspose2d(latent_dim, nb_channels, kernel_size=4),
131             nn.ReLU(inplace=True),
132             nn.ConvTranspose2d(
133                 nb_channels, nb_channels, kernel_size=3, stride=2
134             ),  # from 4x4
135             nn.ReLU(inplace=True),
136             nn.ConvTranspose2d(
137                 nb_channels, nb_channels, kernel_size=4, stride=2
138             ),  # from 9x9
139             nn.ReLU(inplace=True),
140             nn.ConvTranspose2d(nb_channels, nb_channels, kernel_size=5),  # from 20x20
141             nn.ReLU(inplace=True),
142             nn.ConvTranspose2d(nb_channels, 2, kernel_size=5),  # from 24x24
143         )
144
145     def forward(self, z):
146         output = self.model(z.view(z.size(0), -1, 1, 1))
147         mu, log_var = output[:, 0:1], output[:, 1:2]
148         if not args.no_hacks:
149             log_var[...] = 0
150         return mu, log_var
151
152
153 ######################################################################
154
155 data_dir = os.path.join(args.data_dir, "mnist")
156
157 train_set = torchvision.datasets.MNIST(data_dir, train=True, download=True)
158 train_input = train_set.data.view(-1, 1, 28, 28).float()
159
160 test_set = torchvision.datasets.MNIST(data_dir, train=False, download=True)
161 test_input = test_set.data.view(-1, 1, 28, 28).float()
162
163 ######################################################################
164
165 model_q_Z_given_x = LatentGivenImageNet(
166     nb_channels=args.nb_channels, latent_dim=args.latent_dim
167 )
168
169 model_p_X_given_z = ImageGivenLatentNet(
170     nb_channels=args.nb_channels, latent_dim=args.latent_dim
171 )
172
173 optimizer = optim.Adam(
174     itertools.chain(model_p_X_given_z.parameters(), model_q_Z_given_x.parameters()),
175     lr=args.learning_rate,
176 )
177
178 model_p_X_given_z.to(device)
179 model_q_Z_given_x.to(device)
180
181 ######################################################################
182
183 train_input, test_input = train_input.to(device), test_input.to(device)
184
185 train_mu, train_std = train_input.mean(), train_input.std()
186 train_input.sub_(train_mu).div_(train_std)
187 test_input.sub_(train_mu).div_(train_std)
188
189 ######################################################################
190
191 mean_p_Z = train_input.new_zeros(1, args.latent_dim)
192 log_var_p_Z = mean_p_Z
193
194 for epoch in range(args.nb_epochs):
195     acc_loss = 0
196
197     for x in train_input.split(args.batch_size):
198         mean_q_Z_given_x, log_var_q_Z_given_x = model_q_Z_given_x(x)
199         z = sample_gaussian(mean_q_Z_given_x, log_var_q_Z_given_x)
200         mean_p_X_given_z, log_var_p_X_given_z = model_p_X_given_z(z)
201
202         if args.no_dkl:
203             log_q_z_given_x = log_p_gaussian(z, mean_q_Z_given_x, log_var_q_Z_given_x)
204             log_p_x_z = log_p_gaussian(
205                 x, mean_p_X_given_z, log_var_p_X_given_z
206             ) + log_p_gaussian(z, mean_p_Z, log_var_p_Z)
207             loss = -(log_p_x_z - log_q_z_given_x).mean()
208         else:
209             log_p_x_given_z = log_p_gaussian(x, mean_p_X_given_z, log_var_p_X_given_z)
210             dkl_q_Z_given_x_from_p_Z = dkl_gaussians(
211                 mean_q_Z_given_x, log_var_q_Z_given_x, mean_p_Z, log_var_p_Z
212             )
213             loss = (-log_p_x_given_z + dkl_q_Z_given_x_from_p_Z).mean()
214
215         optimizer.zero_grad()
216         loss.backward()
217         optimizer.step()
218
219         acc_loss += loss.item() * x.size(0)
220
221     log_string(f"acc_loss {epoch} {acc_loss/train_input.size(0)}")
222
223 ######################################################################
224
225
226 def save_image(x, filename):
227     x = x * train_std + train_mu
228     x = x.clamp(min=0, max=255) / 255
229     torchvision.utils.save_image(1 - x, filename, nrow=16, pad_value=0.8)
230
231
232 # Save a bunch of test images
233
234 x = test_input[:256]
235 save_image(x, "input.png")
236
237 # Save the same images after encoding / decoding
238
239 mean_q_Z_given_x, log_var_q_Z_given_x = model_q_Z_given_x(x)
240 z = sample_gaussian(mean_q_Z_given_x, log_var_q_Z_given_x)
241 mean_p_X_given_z, log_var_p_X_given_z = model_p_X_given_z(z)
242 if args.no_hacks:
243     x = sample_gaussian(mean_p_X_given_z, log_var_p_X_given_z)
244 else:
245     x = mean_p_X_given_z
246 save_image(x, "output.png")
247
248 # Generate a bunch of images
249
250 z = sample_gaussian(mean_p_Z.expand(x.size(0), -1), log_var_p_Z.expand(x.size(0), -1))
251 mean_p_X_given_z, log_var_p_X_given_z = model_p_X_given_z(z)
252 if args.no_hacks:
253     x = sample_gaussian(mean_p_X_given_z, log_var_p_X_given_z)
254 else:
255     x = mean_p_X_given_z
256 save_image(x, "synth.png")
257
258 ######################################################################