-
Notifications
You must be signed in to change notification settings - Fork 2
/
train_mixer_supervised.py
1450 lines (1325 loc) · 51.6 KB
/
train_mixer_supervised.py
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
428
429
430
431
432
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
470
471
472
473
474
475
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
501
502
503
504
505
506
507
508
509
510
511
512
513
514
515
516
517
518
519
520
521
522
523
524
525
526
527
528
529
530
531
532
533
534
535
536
537
538
539
540
541
542
543
544
545
546
547
548
549
550
551
552
553
554
555
556
557
558
559
560
561
562
563
564
565
566
567
568
569
570
571
572
573
574
575
576
577
578
579
580
581
582
583
584
585
586
587
588
589
590
591
592
593
594
595
596
597
598
599
600
601
602
603
604
605
606
607
608
609
610
611
612
613
614
615
616
617
618
619
620
621
622
623
624
625
626
627
628
629
630
631
632
633
634
635
636
637
638
639
640
641
642
643
644
645
646
647
648
649
650
651
652
653
654
655
656
657
658
659
660
661
662
663
664
665
666
667
668
669
670
671
672
673
674
675
676
677
678
679
680
681
682
683
684
685
686
687
688
689
690
691
692
693
694
695
696
697
698
699
700
701
702
703
704
705
706
707
708
709
710
711
712
713
714
715
716
717
718
719
720
721
722
723
724
725
726
727
728
729
730
731
732
733
734
735
736
737
738
739
740
741
742
743
744
745
746
747
748
749
750
751
752
753
754
755
756
757
758
759
760
761
762
763
764
765
766
767
768
769
770
771
772
773
774
775
776
777
778
779
780
781
782
783
784
785
786
787
788
789
790
791
792
793
794
795
796
797
798
799
800
801
802
803
804
805
806
807
808
809
810
811
812
813
814
815
816
817
818
819
820
821
822
823
824
825
826
827
828
829
830
831
832
833
834
835
836
837
838
839
840
841
842
843
844
845
846
847
848
849
850
851
852
853
854
855
856
857
858
859
860
861
862
863
864
865
866
867
868
869
870
871
872
873
874
875
876
877
878
879
880
881
882
883
884
885
886
887
888
889
890
891
892
893
894
895
896
897
898
899
900
901
902
903
904
905
906
907
908
909
910
911
912
913
914
915
916
917
918
919
920
921
922
923
924
925
926
927
928
929
930
931
932
933
934
935
936
937
938
939
940
941
942
943
944
945
946
947
948
949
950
951
952
953
954
955
956
957
958
959
960
961
962
963
964
965
966
967
968
969
970
971
972
973
974
975
976
977
978
979
980
981
982
983
984
985
986
987
988
989
990
991
992
993
994
995
996
997
998
999
1000
# coding=utf-8
# Copyright 2022 The Google Research Authors.
#
# Licensed under the Apache License, Version 2.0 (the "License");
# you may not use this file except in compliance with the License.
# You may obtain a copy of the License at
#
# http://www.apache.org/licenses/LICENSE-2.0
#
# Unless required by applicable law or agreed to in writing, software
# distributed under the License is distributed on an "AS IS" BASIS,
# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
# See the License for the specific language governing permissions and
# limitations under the License.
# pylint: skip-file
"""A basic CIFAR example using Numpy and JAX.
The primary aim here is simplicity and minimal dependencies.
"""
import jax
import math
import numpy as np
import os
import pickle as pkl
import tensorflow as tf
import time
import jax.numpy as jnp
import optax
import tensorflow_datasets as tfds
from jax.scipy.special import logsumexp
from mixer_lib import (
fa_group_linear,
fa_linear,
get_blk,
get_dataset_metadata,
get_layer_sizes,
get_param_scale,
preprocess,
normalize,
init_random_params,
linear,
group_linear,
depthwise_conv,
get_num_layers,
get_blk_idx,
avg_pooling,
max_pooling,
)
from train_utils import save_checkpoint, last_checkpoint
# Ask TF to not occupy GPU memory.
gpus = tf.config.experimental.list_physical_devices("GPU")
for gpu in gpus:
tf.config.experimental.set_memory_growth(gpu, True)
import augmentations
from spatial_avg_linear import (
spatial_avg_group_linear_custom_jvp,
spatial_avg_group_linear_custom_vjp,
)
from spatial_avg_linear_v2 import (
spatial_avg_group_linear_cross_entropy_custom_jvp,
spatial_avg_group_linear_cross_entropy_custom_vjp,
)
from dataset import _decode_and_random_crop, _decode_and_center_crop
from absl import app
from absl import flags
from jax import grad
from clu import metric_writers
# jax.config.update("jax_array", True)
flags.DEFINE_string("data_root", "datasets", "where to store datasets")
flags.DEFINE_string("exp", "all", "which experiment to run")
flags.DEFINE_string("workdir", "experiments", "experiment directory")
flags.DEFINE_float("mom", 0.0, "momentum")
flags.DEFINE_string("lr", "0.1", "learning rate")
flags.DEFINE_string("optimizer", "sgd", "optimizer name")
flags.DEFINE_float("wd", 0.0001, "weight decay")
flags.DEFINE_integer("warmup_epochs", 1, "number of warmup epochs")
flags.DEFINE_integer("batch_size", 100, "batch size")
flags.DEFINE_integer("num_epochs", 200, "number of epochs")
flags.DEFINE_string("init_scheme", "constant", "kaiming or lecun or constant or list")
flags.DEFINE_integer("num_blocks", 4, "number of mixer blocks")
flags.DEFINE_integer("num_channel_mlp_units", 256, "number of units")
flags.DEFINE_integer("num_channel_mlp_hidden_units", -1, "number of hidden units")
flags.DEFINE_integer("num_token_mlp_units", 64, "number of units")
flags.DEFINE_integer("num_patches", 4, "number of patches on each side")
flags.DEFINE_string("downsample", "1,1,1,1", "downsample ratio")
flags.DEFINE_string("channel_ratio", "1,1,1,1", "channel ratio")
flags.DEFINE_string("group_ratio", "1,1,1,1", "group ratio")
flags.DEFINE_integer("num_channel_group", 1, "number of channels for a group")
flags.DEFINE_integer("num_groups", 1, "number of groups")
flags.DEFINE_float("last_layer_lr", 0.1, "last layer learning rate")
flags.DEFINE_integer("num_passes", 1, "number of forward gradient passes")
flags.DEFINE_bool("layer_norm_all", False, "entire layer")
flags.DEFINE_integer("stop_every", 1, "stop gradient every")
flags.DEFINE_integer("stop_remainder", -1, "stop gradient remainder")
flags.DEFINE_string("schedule", "linear", "learning rate schedule")
flags.DEFINE_bool("aug", False, "whether to do random cropping")
flags.DEFINE_bool("batch_norm", False, "if we run batch norm")
flags.DEFINE_bool("avgpool_token", True, "average pool tokens before head")
flags.DEFINE_bool("concat_groups", True, "whether to concat groups")
flags.DEFINE_float("head_lr", 1.0, "classifier head lr multiplier")
flags.DEFINE_bool("norm_grad", True, "normalization layer gradients")
flags.DEFINE_bool("same_head", True, "whether use same head")
flags.DEFINE_bool("fuse_cross_entropy", True, "whether to fuse operation")
flags.DEFINE_string("dataset", "cifar-10", "cifar-10 or imagenet-100")
flags.DEFINE_bool("begin_ln", False, "layer norm in the beginning")
flags.DEFINE_bool("middle_ln", False, "layer norm in the middle")
flags.DEFINE_bool("last_layer_ln", False, "layer norm before read out")
flags.DEFINE_bool("post_linear_ln", True, "whether to layer norm post linear")
flags.DEFINE_bool("inter_ln", False, "layer norm before intermediate read out")
flags.DEFINE_bool("augcolor", False, "whether to augment colors")
flags.DEFINE_float("area_lb", 0.08, "area crop lower bound")
flags.DEFINE_bool("conv_mixer", False, "use conv for mixing")
flags.DEFINE_integer("kernel_size", 3, "conv kernel size")
flags.DEFINE_bool("bp_last_lr", True, "whether to tune last layer LR for BP")
flags.DEFINE_bool("stopgrad_input", False, "whether to stopgrad on the input")
flags.DEFINE_bool("freeze_backbone", False, "whether freeze")
flags.DEFINE_bool("spatial_loss", True, "whether to keep the spatial dimensions")
flags.DEFINE_bool("modular_loss", True, "whether to have loss functions in each module")
flags.DEFINE_string("gcs_path", "gs://research-brain-rmy-gcp-xgcp", "cloud storage")
flags.DEFINE_bool("use_gcs", False, "whether to use cloud storage")
flags.DEFINE_float("perturb_epsilon", 1e-3, "perturbation stepsize")
flags.DEFINE_bool("train_eval", False, "whether to add train eval")
flags.DEFINE_bool("spatial_pooling", False, "whether to add spatial pooling")
flags.DEFINE_string("pool_fn", "avg", "avg or max")
FLAGS = flags.FLAGS
def normalize_layer(x, swap=False):
return normalize(
x, swap=swap, batch_norm=FLAGS.batch_norm, layer_norm_all=FLAGS.layer_norm_all
)
def mlp_block(
inputs, params, num_groups, noise=None, name="", mask=None, stop_every_layer=False
):
# Token mixing.
print("inputs", inputs.shape)
conv = FLAGS.conv_mixer
states = {}
# B = batch_size, P = num_patches * num_patches
B, P = inputs.shape[0], inputs.shape[1]
G = num_groups
H = int(math.sqrt(P)) # num_patches
if FLAGS.stopgrad_input:
inputs_ = jax.lax.stop_gradient(inputs)
else:
inputs_ = inputs
outputs = jnp.reshape(inputs_, [B, P, -1])
if FLAGS.begin_ln:
outputs = normalize_layer(outputs)
if conv:
outputs = jnp.reshape(outputs, [B, H, H, -1])
states[f"{name}/pre_0"] = outputs
if len(params[0]) == 3: # Feedback alignment
outputs = depthwise_conv(outputs, params[0][0])
outputs_bw = depthwise_conv(outputs, params[0][2])
outputs = (
jax.lax.stop_gradient(outputs - outputs_bw) + outputs_bw + params[0][1]
)
else:
outputs = depthwise_conv(outputs, params[0][0]) + params[0][1]
outputs = jnp.reshape(outputs, [B, P, -1])
else:
outputs = jnp.swapaxes(outputs, 1, 2)
states[f"{name}/pre_0"] = outputs
if len(params[0]) == 3: # Feedback alignment
outputs = fa_linear(outputs, params[0][0], params[0][1], params[0][2])
else:
outputs = linear(outputs, params[0][0], params[0][1])
states[f"{name}/prenorm_0"] = outputs
if FLAGS.post_linear_ln:
outputs = normalize_layer(outputs, swap=not conv)
if noise is not None:
outputs = outputs + noise[0]
outputs = jax.nn.relu(outputs)
if mask is not None:
outputs = outputs * jnp.swapaxes(mask, 1, 2)
states[f"{name}/post_0"] = outputs
if stop_every_layer:
outputs = jax.lax.stop_gradient(outputs)
if not conv:
outputs = jnp.swapaxes(outputs, 1, 2)
# Channel mixing.
if FLAGS.middle_ln:
outputs = normalize(outputs)
states[f"{name}/pre_1"] = outputs
if len(params[1]) == 3: # Feedback alignment
outputs = fa_linear(outputs, params[1][0], params[1][1], params[1][2])
else:
outputs = linear(outputs, params[1][0], params[1][1]) # wx + b
states[f"{name}/prenorm_1"] = outputs
if FLAGS.post_linear_ln:
outputs = normalize_layer(outputs)
if noise is not None:
outputs = outputs + noise[1]
outputs = jax.nn.relu(outputs)
if mask is not None:
outputs = outputs * mask
states[f"{name}/post_1"] = outputs
if stop_every_layer:
outputs = jax.lax.stop_gradient(outputs)
outputs = jnp.reshape(outputs, [B, P, G, -1])
if FLAGS.middle_ln:
outputs = normalize_layer(outputs)
states[f"{name}/pre_2"] = outputs
if len(params[2]) == 3: # Feedback alignment
outputs = fa_group_linear(outputs, params[2][0], params[2][1], params[2][2])
else:
outputs = group_linear(outputs, params[2][0], params[2][1])
states[f"{name}/prenorm_2"] = outputs
if FLAGS.post_linear_ln:
outputs = normalize_layer(outputs)
if noise is not None:
outputs = outputs + noise[2]
if params[1][0].shape[0] != params[1][0].shape[1]:
# Double the channels.
inputs = jnp.concatenate([inputs, inputs], axis=2)
outputs = outputs + inputs # Residual connection
outputs = jax.nn.relu(outputs)
if mask is not None:
outputs = outputs * mask
states[f"{name}/post_2"] = outputs
if stop_every_layer:
outputs = jax.lax.stop_gradient(outputs)
return outputs, (states, {})
def block0(
inputs, params, num_groups, noise=None, name="", mask=None, stop_every_layer=False
):
states = {}
outputs = inputs
if FLAGS.begin_ln:
outputs = normalize_layer(outputs)
states[f"{name}/pre_0"] = outputs
if len(params[0]) == 3: # Feedback alignment
outputs = fa_linear(outputs, params[0][0], params[0][1], params[0][2])
else:
outputs = linear(
outputs, params[0][0], params[0][1]
) # param[0][0] = W, param[0][1] = b
states[f"{name}/prenorm_0"] = outputs
if FLAGS.post_linear_ln:
outputs = normalize_layer(outputs)
if noise is not None:
outputs = outputs + noise[0]
outputs = jax.nn.relu(outputs)
if mask is not None:
outputs = outputs * mask
states[f"{name}/post_0"] = outputs
if stop_every_layer:
outputs = jax.lax.stop_gradient(outputs)
B, P, D = outputs.shape # B = batch size, P = patch size, D = embedding dimension
G = num_groups
outputs = jnp.reshape(outputs, [B, P, G, -1])
if FLAGS.middle_ln:
outputs = normalize_layer(outputs)
states[f"{name}/pre_1"] = outputs
if len(params[1]) == 3: # Feedback alignment
outputs = fa_group_linear(outputs, params[1][0], params[1][1], params[1][2])
else:
outputs = group_linear(outputs, params[1][0], params[1][1])
states[f"{name}/prenorm_1"] = outputs
if FLAGS.post_linear_ln:
outputs = normalize_layer(outputs)
if noise is not None:
outputs = outputs + noise[1]
outputs = jax.nn.relu(outputs)
if mask is not None:
outputs = outputs * mask
states[f"{name}/post_1"] = outputs
if stop_every_layer:
outputs = jax.lax.stop_gradient(outputs)
return outputs, (states, {})
def run_block(
block_idx, num_groups, inputs, block_params, block_noise, stop_every_layer=False
):
if block_idx == 0:
outputs, (states, logs) = block0(
inputs,
block_params,
num_groups,
block_noise,
name=f"block_{block_idx}",
stop_every_layer=stop_every_layer,
)
elif block_idx > 0:
outputs, (states, logs) = mlp_block(
inputs,
block_params,
num_groups,
block_noise,
name=f"block_{block_idx}",
stop_every_layer=stop_every_layer,
)
else:
# Negative number is the final block.
states = {}
x = jnp.reshape(inputs, [inputs.shape[0], inputs.shape[1], -1])
x = jnp.mean(x, axis=1) # [B, D]
# For supervised classification readout (unsupervised rep learning).
x = jax.lax.stop_gradient(x)
states[f"pre_cls"] = x
pred_cls = linear(x, block_params[-1][0], block_params[-1][1])
states["pred_cls"] = pred_cls
outputs = pred_cls
logs = {}
return outputs, (states, logs)
def predict(
params,
inputs,
noise=None,
stop_gradient=False,
readout=False,
stop_every=1,
stop_remainder=-1,
is_training=False,
stop_every_layer=False,
):
"""MLP mixer"""
NBLK = FLAGS.num_blocks
downsample = [int(d) for d in FLAGS.downsample.split(",")]
group_ratio = [int(d) for d in FLAGS.group_ratio.split(",")]
if stop_remainder < 0:
stop_remainder = stop_remainder + stop_every
md = get_dataset_metadata(FLAGS.dataset)
inputs = preprocess(inputs, md["image_mean"], md["image_std"], FLAGS.num_patches)
x = inputs
# We will start with a channel mixing MLP instead of token mixing.
all_states = {}
all_logs = {}
num_groups_ = FLAGS.num_groups
# Build network.
for blk in range(NBLK):
start, end = get_blk_idx(blk)
if noise is not None:
noise_ = noise[start:end]
else:
noise_ = None
x, (states, logs) = run_block(
blk,
num_groups_,
x,
params[start:end],
noise_,
stop_every_layer=stop_every_layer,
)
x_proj = x
if FLAGS.inter_ln: # layer norm before intermedite read out
x_proj = normalize(x_proj)
states[f"block_{blk}/pre_pred"] = x_proj
# all_states = all_states | states
for k in states:
all_states[k] = states[k]
# all_logs = all_logs | logs
for k in logs:
all_logs[k] = logs[k]
# FLAGS.stopgrad_input: whether to stopgrad on the input
if stop_gradient and not FLAGS.stopgrad_input:
if blk % stop_every == stop_remainder:
x = jax.lax.stop_gradient(x)
if downsample[blk] > 1:
# Downsample 2x
if FLAGS.pool_fn == "avg":
x = avg_pooling(x, stride=downsample[blk])
elif FLAGS.pool_fn == "max":
x = max_pooling(x, stride=downsample[blk])
num_groups_ = num_groups_ * group_ratio[blk]
if readout:
x = jax.lax.stop_gradient(x)
x = jnp.reshape(x, [x.shape[0], x.shape[1], -1])
x = jnp.mean(x, axis=1) # [B, D]
if FLAGS.last_layer_ln:
x = normalize(x)
all_states[f"pre_final"] = x
# [B, K]
if len(params[-1]) == 3:
pred = fa_linear(x, params[-1][0], params[-1][1], params[-1][2])
else:
pred = linear(x, params[-1][0], params[-1][1])
all_states["pred_final"] = pred
return pred, (all_states, all_logs)
def loss_dfa(params, batch, noise=None, key=None):
if FLAGS.augcolor:
key, subkey = jax.random.split(key)
batch = augmentations.postprocess1(batch, subkey, add_gaussian_blur=False)
inputs, targets = batch["image"], batch["label"]
md = get_dataset_metadata(FLAGS.dataset)
targets_onehot = jax.nn.one_hot(targets, md["num_classes"])
logits, (states, logs) = predict(
params,
inputs,
noise=noise,
stop_gradient=True,
stop_every_layer=True,
readout=True,
is_training=True,
)
loss = jnp.mean(classif_loss(logits, targets_onehot))
predicted_class = jnp.argmax(logits, axis=-1)
logs["acc/train"] = jnp.mean(predicted_class == targets)
logs["loss"] = loss
local_losses = []
NBLK = FLAGS.num_blocks
NL = get_num_layers(NBLK)
for j, (weights, bias, weights_b, bias_b) in enumerate(params[:NL]):
blk, layer = get_blk(j)
x_proj = states[f"block_{blk}/post_{layer}"]
B, P = x_proj.shape[0], x_proj.shape[1]
x_proj = jnp.mean(x_proj, [1])
x_proj = jnp.reshape(x_proj, [B, -1])
logit_bw_ = linear(x_proj, weights_b, bias_b)
local_loss_bw_ = jnp.mean(classif_loss(logit_bw_, targets_onehot))
local_loss_ = jax.lax.stop_gradient(loss - local_loss_bw_) + local_loss_bw_
local_losses.append(local_loss_)
logs[f"local_loss/blk_{blk}"] = local_loss_
return loss, local_losses, (states, logs)
def loss(
params,
batch,
noise=None,
stop_gradient=False,
readout=False,
key=None,
stop_every=1,
stop_remainder=-1,
custom_forward=False,
avg_batch=True,
):
"""Supervised classification loss."""
if FLAGS.augcolor:
key, subkey = jax.random.split(key)
batch = augmentations.postprocess1(batch, subkey, add_gaussian_blur=False)
inputs, targets = batch["image"], batch["label"]
md = get_dataset_metadata(FLAGS.dataset)
targets_onehot = jax.nn.one_hot(targets, md["num_classes"])
logits, (states, logs) = predict(
params,
inputs,
noise=noise,
stop_gradient=stop_gradient,
readout=readout,
stop_every=stop_every,
stop_remainder=stop_remainder,
is_training=True,
)
loss = classif_loss(logits, targets_onehot) # cross entropy loss
if avg_batch:
loss = jnp.mean(loss)
predicted_class = jnp.argmax(logits, axis=-1)
logs["acc/train"] = jnp.mean(predicted_class == targets)
if not avg_batch:
logs["loss"] = jnp.mean(loss)
else:
logs["loss"] = loss
local_losses = []
NBLK = FLAGS.num_blocks
NL = get_num_layers(NBLK)
avgpool_token = FLAGS.avgpool_token
for blk in range(NBLK):
x_proj = states[f"block_{blk}/pre_pred"]
B, P = x_proj.shape[0], x_proj.shape[1] # B: batch size, P: number of patches
param_ = params[NL + blk]
if FLAGS.fuse_cross_entropy:
if custom_forward:
loss_fn_ = spatial_avg_group_linear_cross_entropy_custom_jvp
else:
loss_fn_ = spatial_avg_group_linear_cross_entropy_custom_vjp
local_loss_ = loss_fn_(x_proj, param_[0], param_[1], targets_onehot)
if len(param_) == 3: # Feedback alignment
local_loss_bw_ = loss_fn_(x_proj, param_[2], param_[1], targets_onehot)
local_loss_ = (
jax.lax.stop_gradient(local_loss_ - local_loss_bw_) + local_loss_bw_
)
else:
assert False
if not FLAGS.spatial_loss:
# [B, P, G] -> [B, 1, G]
local_loss_ = jnp.mean(local_loss_, axis=1, keepdims=True)
assert avgpool_token
if avgpool_token:
denom = B
else:
denom = B * P
local_losses.append(local_loss_ / float(denom))
logs[f"local_loss/blk_{blk}"] = jnp.mean(local_loss_)
if not FLAGS.modular_loss:
local_losses = local_losses[-1]
return loss, local_losses, (states, logs)
def classif_loss(logits, targets):
"""
Calculate the cross-entropy loss between the logits and the labels.
"""
logits = logits - logsumexp(logits, axis=-1, keepdims=True)
if len(logits.shape) == 3:
targets = targets[:, None, :]
elif len(logits.shape) == 4:
targets = targets[:, None, None, :]
elif len(logits.shape) == 5:
targets = targets[:, None, None, None, :]
loss = -jnp.sum(logits * targets, axis=-1)
return loss
def accuracy(params, batch):
inputs, targets = batch["image"], batch["label"]
pred, (states, logs) = predict(params, inputs)
predicted_class = jnp.argmax(pred, axis=-1)
return jnp.mean(predicted_class == targets)
def update_backprop(params, batch, key):
wd = FLAGS.wd
def loss_fn(params, batch):
final_loss_, local_loss_, (states, logs) = loss(
params, batch, noise=None, stop_gradient=False, readout=False, key=key
)
return final_loss_, (states, logs)
grads_now, (states, logs) = grad(loss_fn, has_aux=True)(params, batch)
if FLAGS.freeze_backbone:
grads_now = [(0.0 * gw, 0.0 * gb) for (gw, gb) in grads_now[:-1]] + [
grads_now[-1]
]
if FLAGS.last_layer_lr < 1.0 and FLAGS.bp_last_lr:
grads_now[-1] = (
grads_now[-1][0] * FLAGS.last_layer_lr,
grads_now[-1][1] * FLAGS.last_layer_lr,
)
if FLAGS.optimizer == "sgd" and wd > 0.0:
grads_now = [(gw + wd * w, gb) for (gw, gb), (w, b) in zip(grads_now, params)]
return grads_now, logs
def update_forward_grad_weights(params, batch, key):
num_patches = FLAGS.num_patches
num_groups = FLAGS.num_groups
num_passes = FLAGS.num_passes
md = get_dataset_metadata(FLAGS.dataset)
grads_now = [
(jnp.zeros_like(weight), jnp.zeros_like(bias)) for (weight, bias) in params
]
if FLAGS.fuse_cross_entropy:
local_classif_loss = lambda x, w, b, targets: jnp.mean(
spatial_avg_group_linear_cross_entropy_custom_vjp(x, w, b, targets)
)
else:
local_classif_loss = lambda x, w, b, targets: jnp.mean(
classif_loss(spatial_avg_group_linear_custom_vjp(x, w, b), targets)
)
local_classif_grad = jax.grad(local_classif_loss, argnums=[1, 2])
global_classif_loss = lambda x, w, b, targets: jnp.mean(
classif_loss(jnp.einsum("nc,cd->nd", x, w) + b, targets)
)
global_classif_grad = jax.grad(global_classif_loss, argnums=[1, 2])
wd = FLAGS.wd
def local_loss(params):
_, local_loss_, (states, logs) = loss(
params,
batch,
stop_gradient=FLAGS.modular_loss,
readout=True,
key=key,
custom_forward=True,
)
return local_loss_, (states, logs)
for npass in range(num_passes):
noise = []
num_items = num_patches**2
NBLK = FLAGS.num_blocks
label = jax.nn.one_hot(batch["label"], md["num_classes"])
G = num_groups
M = num_passes
NL = get_num_layers(NBLK) # number of layers in the main network
main_params = params[:NL] # 11 of the 16 layers (main network layers)
loss_params = params[NL:] # 5 of the 16 layers (classification layers)
for i, (weight, bias) in enumerate(main_params): # add noise to main_params
# blk, layer = get_blk(i) # useless code
key, subkey = jax.random.split(key)
dw = jax.random.normal(subkey, weight.shape)
key, subkey = jax.random.split(key)
db = jax.random.normal(subkey, bias.shape)
noise.append((dw, db))
for i, (weight, bias) in enumerate(loss_params): # add noise to loss_params
noise.append((jnp.zeros_like(weight), jnp.zeros_like(bias)))
# [L,B,P]
_, g, (states, logs) = jax.jvp(local_loss, [params], [noise], has_aux=True)
# Main network layers.
for i, ((weight, bias), (dw, db)) in enumerate(zip(main_params, noise)):
blk, layer = get_blk(i)
# Forward gradient.
if FLAGS.modular_loss:
g_ = g[blk]
else:
g_ = g
# [B, P, G] -> [G]
g_ = jnp.sum(g_, axis=[0, 1])
# 1st layer of blocks following the first block
if blk > 0 and layer == 0:
g_ = jnp.sum(g_) # []
grad_w = g_ * dw
grad_b = g_ * db
# 1st layer of first block or 2nd layer of following blocks
elif (blk == 0 and layer == 0) or (blk > 0 and layer == 1):
dw = jnp.reshape(dw, [weight.shape[0], G, weight.shape[1] // G])
db = jnp.reshape(db, [G, -1])
grad_w = g_[None, :, None] * dw
grad_b = g_[:, None] * db
# 2nd layer of first block or 3rd layer of following blocks
elif (blk == 0 and layer == 1) or layer == 2:
dw = jnp.reshape(dw, [G, weight.shape[0], -1])
db = jnp.reshape(db, [G, -1])
grad_w = g_[:, None, None] * dw
grad_b = g_[:, None] * db
print(blk, layer, weight.shape, grad_w.shape)
# assert False
grad_w = jnp.reshape(grad_w, weight.shape)
grad_b = jnp.reshape(grad_b, bias.shape)
idx = i
grads_now[idx] = (
grads_now[idx][0] + grad_w / float(M),
grads_now[idx][1] + grad_b / float(M),
)
# Classification layers.
for i, (weight, bias) in enumerate(loss_params):
blk = i # Every block has a loss.
if i == len(loss_params) - 1:
# Last classification layer.
pre_act = states[f"pre_final"]
grad_w, grad_b = global_classif_grad(pre_act, weight, bias, label)
grad_w = grad_w * FLAGS.last_layer_lr
grad_b = grad_b * FLAGS.last_layer_lr
else:
# Intermediate classification layer.
pre_act = states[f"block_{blk}/pre_pred"]
grad_w, grad_b = local_classif_grad(pre_act, weight, bias, label)
grad_w = grad_w * FLAGS.head_lr
grad_b = grad_b * FLAGS.head_lr
idx = i + NL
grads_now[idx] = (
grads_now[idx][0] + grad_w / float(M),
grads_now[idx][1] + grad_b / float(M),
)
if FLAGS.optimizer == "sgd" and wd > 0.0:
grads_now = [(gw + wd * w, gb) for (gw, gb), (w, b) in zip(grads_now, params)]
return grads_now, logs
def sample_activation_noise(
params,
key,
batch_size,
num_patches,
num_hid_units,
num_units,
num_groups,
num_blocks,
conv=False,
num_passes=1,
):
noise = []
H = num_hid_units
D = num_units
G = num_groups
P = num_patches
group_ratio = [int(g) for g in FLAGS.group_ratio.split(",")]
downsample = [int(d) for d in FLAGS.downsample.split(",")]
channel_ratio = [int(c) for c in FLAGS.channel_ratio.split(",")]
B = batch_size
NBLK = num_blocks
NL = get_num_layers(NBLK)
main_params = params[:NL]
loss_params = params[NL:]
group_list = [G]
token_list = [P]
hid_unit_list = [H]
unit_list = [D]
for blk in range(NBLK - 1):
group_list.append(group_list[-1] * group_ratio[blk])
token_list.append(token_list[-1] // (downsample[blk] ** 2))
hid_unit_list.append(hid_unit_list[-1] * channel_ratio[blk])
unit_list.append(unit_list[-1] * channel_ratio[blk])
for i, p in enumerate(main_params):
blk, layer = get_blk(i)
G_ = group_list[blk]
P_ = token_list[blk]
H_ = hid_unit_list[blk]
if blk > 0:
I_ = unit_list[blk - 1]
D_ = unit_list[blk]
if blk == 0:
if layer == 0:
out_shape = [B, P_, H_]
else:
out_shape = [B, P_, G_, D_ // G_]
else:
if layer == 0:
if conv:
out_shape = [B, P_, I_]
else:
out_shape = [B, I_, P_]
elif layer == 1:
out_shape = [B, P_, H_]
else:
out_shape = [B, P_, G_, D_ // G_]
if num_passes > 1:
out_shape = [num_passes] + out_shape
key, subkey = jax.random.split(key)
dz = jax.random.normal(subkey, out_shape)
noise.append(dz)
return noise, key
def update_forward_grad_activations(params, batch, key):
num_patches = FLAGS.num_patches
num_units = FLAGS.num_channel_mlp_units
num_hid_units = FLAGS.num_channel_mlp_hidden_units
num_items = num_patches**2
NBLK = FLAGS.num_blocks
if num_hid_units < 0:
num_hid_units = num_units
num_groups = FLAGS.num_groups
group_ratio = [int(g) for g in FLAGS.group_ratio.split(",")]
downsample = [int(d) for d in FLAGS.downsample.split(",")]
channel_ratio = [int(c) for c in FLAGS.channel_ratio.split(",")]
group_list = [num_groups]
token_list = [num_patches**2]
hid_unit_list = [num_hid_units]
unit_list = [num_units]
for blk in range(NBLK - 1):
group_list.append(group_list[-1] * group_ratio[blk])
token_list.append(token_list[-1] // (downsample[blk] ** 2))
hid_unit_list.append(hid_unit_list[-1] * channel_ratio[blk])
unit_list.append(unit_list[-1] * channel_ratio[blk])
num_passes = FLAGS.num_passes
conv = FLAGS.conv_mixer
md = get_dataset_metadata(FLAGS.dataset)
grads_now = [
(jnp.zeros_like(weight), jnp.zeros_like(bias)) for (weight, bias) in params
]
ln_loss = lambda x, dy: jnp.sum(normalize(x) * dy)
ln_grad = jax.grad(ln_loss)
ln_loss1 = lambda x, dy: jnp.sum(normalize(x, swap=not conv) * dy)
ln_grad1 = jax.grad(ln_loss1)
conv_loss = lambda w, x, dy: jnp.sum(depthwise_conv(x, w) * dy)
conv_grad = jax.grad(conv_loss)
if FLAGS.fuse_cross_entropy:
local_classif_loss = lambda x, w, b, targets: jnp.mean(
spatial_avg_group_linear_cross_entropy_custom_vjp(x, w, b, targets)
)
else:
local_classif_loss = lambda x, w, b, targets: jnp.mean(
classif_loss(spatial_avg_group_linear_custom_vjp(x, w, b), targets)
)
local_classif_grad = jax.grad(local_classif_loss, argnums=[1, 2])
global_classif_loss = lambda x, w, b, targets: jnp.mean(
classif_loss(jnp.einsum("nc,cd->nd", x, w) + b, targets)
)
global_classif_grad = jax.grad(global_classif_loss, argnums=[1, 2])
wd = FLAGS.wd
def local_loss(noise):
_, local_loss_, (states, logs) = loss(
params,
batch,
noise=noise,
stop_gradient=FLAGS.modular_loss,
readout=True,
key=key,
custom_forward=True,
)
return local_loss_, (states, logs)
for npass in range(num_passes):
zeros = []
noise = []
B = batch["image"].shape[0]
label = jax.nn.one_hot(batch["label"], md["num_classes"])
M = num_passes
NL = get_num_layers(NBLK)
main_params = params[:NL]
loss_params = params[NL:]
noise, key = sample_activation_noise(
params,
key,
B,
num_items,
num_hid_units,
num_units,
num_groups,
NBLK,
conv=conv,
num_passes=1,
)
zeros = jax.tree_util.tree_map(lambda x: jnp.zeros_like(x), noise)
# [L,B,P]
_, g, (states, logs) = jax.jvp(local_loss, [zeros], [noise], has_aux=True)
# Main network layers.
for i, ((weight, bias), dz) in enumerate(zip(main_params, noise)):
blk, layer = get_blk(i)
# Forward gradient.
pre_act = states[f"block_{blk}/pre_{layer}"]
prenorm_act = states[f"block_{blk}/prenorm_{layer}"]
post_act = states[f"block_{blk}/post_{layer}"]
mask = (post_act > 0.0).astype(jnp.float32)
G_ = group_list[blk]
P_ = token_list[blk]
if FLAGS.modular_loss:
g_ = g[blk]
else:
g_ = g
# [B, D] -> [B, D, P] or [B, P] -> [B, P, D]
if blk > 0 and layer == 0:
# Token mixing layer
if conv:
g_ = jnp.sum(g_, axis=-1, keepdims=True) # [2B, P, 1]
else:
g_ = jnp.sum(g_, axis=-1)[:, None, :] # [2B, 1, P]
else:
# Channel mixing layer
dz = jnp.reshape(dz, [B, P_, G_, -1])
mask = jnp.reshape(mask, [B, P_, G_, -1])
g_ = g_[:, :, :, None]
grad_z = g_ * dz * mask
# Backprop through normalization.
if FLAGS.norm_grad and FLAGS.post_linear_ln:
if blk > 0 and layer == 0:
grad_z = jnp.reshape(
ln_grad1(prenorm_act, jnp.reshape(grad_z, prenorm_act.shape)),
dz.shape,
)
else:
grad_z = jnp.reshape(
ln_grad(prenorm_act, jnp.reshape(grad_z, prenorm_act.shape)),
dz.shape,
)
if blk > 0 and layer == 0 and conv:
# Token mixing conv
H_ = int(math.sqrt(P_))
grad_z = jnp.reshape(grad_z, [B, H_, H_, -1])
grad_w = conv_grad(weight, pre_act, grad_z)
grad_b = jnp.einsum("nhwd->d", grad_z)
elif blk > 0 and layer == 0 and not conv:
# Token mixing FC
grad_w = jnp.einsum("npc,npd->cd", pre_act, grad_z)
grad_b = jnp.einsum("npc->c", grad_z)
elif (blk == 0 and layer == 0) or (blk > 0 and layer == 1):
grad_z = jnp.reshape(grad_z, [B, P_, -1])
# Channel mixing FC
grad_w = jnp.einsum("npc,npd->cd", pre_act, grad_z)
grad_b = jnp.einsum("npc->c", grad_z)
else:
# Channel mixing group FC
grad_w = jnp.einsum("npgc,npgd->gcd", pre_act, grad_z)
grad_b = jnp.einsum("npgd->gd", grad_z)
idx = i
grads_now[idx] = (
grads_now[idx][0] + grad_w / float(M),
grads_now[idx][1] + grad_b / float(M),
)
# Classification layers.
for i, (weight, bias) in enumerate(loss_params):
blk = i # Every block has a loss.
if i == len(loss_params) - 1:
# Last classification layer.
pre_act = states[f"pre_final"]
grad_w, grad_b = global_classif_grad(pre_act, weight, bias, label)
grad_w = grad_w * FLAGS.last_layer_lr
grad_b = grad_b * FLAGS.last_layer_lr
else:
# Intermediate classification layer.
pre_act = states[f"block_{blk}/pre_pred"]
grad_w, grad_b = local_classif_grad(pre_act, weight, bias, label)
grad_w = grad_w * FLAGS.head_lr
grad_b = grad_b * FLAGS.head_lr
idx = i + NL
grads_now[idx] = (
grads_now[idx][0] + grad_w / float(M),
grads_now[idx][1] + grad_b / float(M),
)
if FLAGS.freeze_backbone:
grads_now = [(0.0 * gw, 0.0 * gb) for (gw, gb) in grads_now[:-1]] + [
grads_now[-1]
]
if FLAGS.optimizer == "sgd" and wd > 0.0:
grads_now = [(gw + wd * w, gb) for (gw, gb), (w, b) in zip(grads_now, params)]
return grads_now, logs
def _to_tfds_split(split):
"""Returns the TFDS split appropriately sharded."""
if split in ["train", "valid", "train_eval"]:
return tfds.Split.TRAIN
elif split == "test":
return tfds.Split.TEST
def get_dataset_cifar10(split, seed=0):
batch_size = FLAGS.batch_size
data_root = FLAGS.data_root
ds = tfds.load(
"cifar10", split=_to_tfds_split(split), data_dir=data_root, shuffle_files=True
)
ds = ds.repeat()
ds = ds.shuffle(buffer_size=10 * batch_size, seed=seed)
is_parallel = jax.device_count() > 1
num_parallel = jax.local_device_count()
def preprocess(example):
image = tf.image.convert_image_dtype(example["image"], tf.float32)
if split == "train":
if FLAGS.aug:
image = tf.pad(image, [(4, 4), (4, 4), (0, 0)])
image = tf.image.random_crop(image, (32, 32, 3))
image = tf.image.random_flip_left_right(image)
elif split == "train_noaug":
image = tf.image.random_flip_left_right(image)
label = tf.cast(example["label"], tf.int32)
return {"image": image, "label": label}
ds = ds.map(preprocess, num_parallel_calls=tf.data.experimental.AUTOTUNE)
ds = ds.batch(batch_size)
def rebatch(batch):
if is_parallel:
for k in batch:
batch[k] = tf.reshape(
batch[k],
[num_parallel, batch_size // num_parallel]
+ list(batch[k].shape[1:]),
)
return batch
ds = ds.map(rebatch, num_parallel_calls=tf.data.experimental.AUTOTUNE)
ds = ds.prefetch(tf.data.experimental.AUTOTUNE)
yield from tfds.as_numpy(ds)
def _to_imagenet100_split(split):