3 # Any copyright is dedicated to the Public Domain.
4 # https://creativecommons.org/publicdomain/zero/1.0/
6 # Written by Francois Fleuret <francois@fleuret.org>
8 # This is an implementation from scratch of a "GPT", that is a model
9 # composed of several causal self-attention blocks. It is equipped
10 # with a caching mechanism for keys and values to avoid a O(N^3) cost
11 # for auto-regression.
13 # This implementation is equipped with RNN layers to replace the MHA
20 from torch.nn import functional as F
26 ######################################################################
28 # A BracketedSequence is a BxTx... tensor with a first and a nb time
31 # Modules able to process it expect that they will have to process a
32 # first bracket starting at t=0, followed by a succession of brackets
33 # that move forward in time, do not overlap, and cover the axis T with
36 # Although it is more general, for a classical prompt-conditioned
37 # auto-regressive process it will be a first bracket starting at 0 and
38 # of arbitrary length for the "prompt", followed by brackets of length
39 # 1 for the successive tokens.
41 # Modules able to process brackets may implement a cache that is
42 # resetted when init_cache is True
45 class BracketedSequence:
46 def __init__(self, x, first=None, nb=None, init_cache=None):
48 assert (first is None and nb is None and init_cache is None) or (
49 first is not None and nb is not None and init_cache is not None
52 self.first = 0 if first is None else first
53 self.nb = x.size(1) if nb is None else nb
54 self.init_cache = True if init_cache is None else init_cache
57 return self.x[:, self.first : self.first + self.nb]
60 return self.first == 0 and self.nb == self.x.size(1)
63 ######################################################################
66 class CacheWrapper(nn.Module):
67 def __init__(self, *f):
69 self.f = f[0] if len(f) == 1 else nn.Sequential(*f)
71 def forward(self, bs):
73 y = self.f(bs.slice())
74 self.cache_y = y.new(*((y.size(0), bs.x.size(1)) + y.size()[2:]))
75 self.cache_y[:, bs.first : bs.first + bs.nb] = y
77 assert tuple(bs.x.size()[:2]) == tuple(self.cache_y.size()[:2])
78 assert bs.first + bs.nb <= self.cache_y.size(1)
79 self.cache_y[:, bs.first : bs.first + bs.nb] = self.f(bs.slice())
81 return BracketedSequence(self.cache_y, bs.first, bs.nb, bs.init_cache)
84 ##############################
87 class WithResidual(nn.Module):
88 def __init__(self, *f):
90 self.f = f[0] if len(f) == 1 else nn.Sequential(*f)
92 def forward(self, bs):
93 return BracketedSequence(bs.x + self.f(bs).x, bs.first, bs.nb, bs.init_cache)
96 ##############################
99 class AddPositionalEncoding(nn.Module):
100 def __init__(self, len_max):
102 self.len_max = len_max
104 # [Vaswani et al 2018] PE_{t,2i} = sin(t/(L^{2i/D})), PE_{t,2i+1} = cos(t/(L^{2i/D}))
106 def forward(self, bs):
108 t = torch.arange(bs.x.size(1), dtype=bs.x.dtype, device=bs.x.device)[
111 j = torch.arange(bs.x.size(2), dtype=bs.x.dtype, device=bs.x.device)[
116 t / (self.len_max ** ((j - k) / bs.x.size(2))) + math.pi / 2 * k
118 self.cache_y = bs.x.new(bs.x.size())
120 self.cache_y[:, bs.first : bs.first + bs.nb] = (
121 bs.slice() + self.pe[bs.first : bs.first + bs.nb]
124 return BracketedSequence(self.cache_y, bs.first, bs.nb, bs.init_cache)
130 # X is /.../xTxD A is /.../xT Y_init is /.../xD
133 def pscan_dim(A, X, Y_init, dim=-2):
135 a, T, b = s[:dim].numel(), s[dim], s[dim + 1 :].numel()
137 A = A.reshape(a, T, *s[dim + 1 : -1])
138 X = X.reshape(a, T, *s[dim + 1 : -1], -1)
141 Y_init = X.new_zeros(a, *s[dim + 1 : -1], X.size(-1))
143 Y_init = Y_init.reshape(a, *s[dim + 1 : -1], -1)
145 Y = pscan.pscan(A, X, Y_init).reshape(s)
150 def pscan_shape(A, X, Y_init):
152 A = A.reshape(-1, s[-2])
153 X = X.reshape(-1, s[-2], s[-1])
156 Y_init = X.new_zeros(X.size(0), s[-1])
158 Y_init = Y_init.reshape(-1, s[-1])
160 Y = pscan.pscan(A, X, Y_init).reshape(s)
165 def nsum_shape(X, Y_init):
167 X = X.reshape(-1, s[-2], s[-1]) # ntd
169 Y = 0 if Y_init is None else Y_init.reshape(-1, s[-1])
172 for k in range(X.size(1)):
174 Y = Y / Y.norm(dim=-1, keepdim=True).clamp(min=1)
177 return torch.cat(result, dim=1).reshape(s)
180 ##############################
183 class DumbRec(nn.Module):
191 attention_dropout=0.0,
199 return nn.Parameter(torch.randn(*d) / math.sqrt(d[-1]))
201 self.nb_lines = nb_lines
202 self.attention_dropout = attention_dropout
204 self.k_star = randw(nb_lines, dim_qk)
206 self.w_qw = randw(nb_heads, dim_qk, dim_model)
207 self.w_qr = randw(nb_heads, dim_qk, dim_model)
208 # self.w_k = randw(nb_heads, dim_qk, dim_model)
209 self.w_v = randw(nb_heads, dim_v, dim_model)
210 self.w_o = randw(dim_v * nb_heads, dim_model)
212 def reset_inner_loss(self):
213 self.acc_attention = 0
216 def get_inner_loss(self):
217 warnings.warn("l2 regularization", RuntimeWarning)
218 return (self.acc_attention / self.acc_nb).pow(2).sum()
219 # return torch.tensor([0], device=self.w_qw.device)
221 def forward(self, bs):
222 x_q, t0, t1 = bs.x, bs.first, bs.first + bs.nb
225 self.rec_v = x_q.new_zeros(
226 x_q.size(0), self.nb_lines, x_q.size(1), self.w_v.size(1)
228 # self.rec_k = x_q.new_zeros(
229 # x_q.size(0), self.nb_lines, x_q.size(1), self.w_k.size(1)
231 self.cache_y = x_q.new_zeros(x_q.size(0), x_q.size(1), self.w_o.size(1))
233 ######################################################################
236 k_star = self.k_star[:, None, :].expand(-1, t1 - t0, -1)
238 warnings.warn("rotating key barrel", RuntimeWarning)
239 k_star = self.k_star[:, None, :].expand(-1, x_q.size(1), -1)
240 t_barrel = torch.arange(t0, t1, device=k_star.device)
241 t_barrel = t_barrel[None, :].expand(k_star.size(0), t1 - t0)
243 torch.arange(k_star.size(0), device=k_star.device)[:, None] + t_barrel
245 k_star = k_star[l_barrel, t_barrel]
247 ######################################################################
248 # Compute the recurrent state
250 qw = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_qw)
252 v = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_v)
253 # k = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_k)
259 ) / math.sqrt(self.w_qw.size(1))
261 aw = aw.softmax(dim=2) # nhlt
264 self.acc_attention += aw.sum(dim=(0, 1, 3))
265 self.acc_nb += aw.size(0) * aw.size(1) * aw.size(3)
267 aw = F.dropout(aw, self.attention_dropout, self.training)
269 A = 1 - aw.sum(dim=1) # nlt
271 V = torch.einsum("nhlt,nhtd->nltd", aw, v).contiguous()
272 # K = torch.einsum("nhlt,nhtd->nltd", aw, k).contiguous()
278 V0 = self.rec_v[:, :, t0 - 1]
279 # K0 = self.rec_k[:, :, t0 - 1]
281 self.rec_v[:, :, t0:t1] = pscan_shape(A, V, V0)
282 # self.rec_k[:, :, t0:t1] = pscan_shape(A, K, K0)
284 ######################################################################
285 # compute the readout
287 qr = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_qr)
292 # self.rec_k[:, :, t0:t1],
294 ) / math.sqrt(self.w_qr.size(1))
296 ar = ar.softmax(dim=2) # nhlt
298 ar = F.dropout(ar, self.attention_dropout, self.training)
303 self.rec_v[:, :, t0:t1],
306 self.cache_y[:, t0:t1] = y @ self.w_o
308 return BracketedSequence(self.cache_y, t0, t1 - t0, bs.init_cache)
311 ##############################
314 class KVRec(nn.Module):
322 attention_dropout=0.0,
330 return nn.Parameter(torch.randn(*d) / math.sqrt(d[-1]))
332 self.nb_lines = nb_lines
333 self.attention_dropout = attention_dropout
335 self.k_star = randw(nb_lines, dim_qk)
337 self.w_qw = randw(nb_heads, dim_qk, dim_model)
338 self.w_qr = randw(nb_heads, dim_qk, dim_model)
339 self.w_k = randw(nb_heads, dim_qk, dim_model)
340 self.w_v = randw(nb_heads, dim_v, dim_model)
341 self.w_o = randw(dim_v * nb_heads, dim_model)
343 def reset_inner_loss(self):
344 self.acc_attention = 0
347 def get_inner_loss(self):
348 warnings.warn("l2 regularization", RuntimeWarning)
349 return (self.acc_attention / self.acc_nb).pow(2).sum()
350 # return torch.tensor([0], device=self.w_qw.device)
351 # warnings.warn("side regularization", RuntimeWarning)
353 # (0.5 / self.nb_lines - self.acc_attention / self.acc_nb).clamp(min=0).sum()
355 # return torch.tensor([0], device=self.w_qw.device)
357 def forward(self, bs):
358 x_q, t0, t1 = bs.x, bs.first, bs.first + bs.nb
361 self.rec_v = x_q.new_zeros(
362 x_q.size(0), self.nb_lines, x_q.size(1), self.w_v.size(1)
364 self.rec_k = x_q.new_zeros(
365 x_q.size(0), self.nb_lines, x_q.size(1), self.w_k.size(1)
367 self.cache_y = x_q.new_zeros(x_q.size(0), x_q.size(1), self.w_o.size(1))
369 ######################################################################
372 k_star = self.k_star[:, None, :].expand(-1, t1 - t0, -1)
374 warnings.warn("rotating key barrel", RuntimeWarning)
375 k_star = self.k_star[:, None, :].expand(-1, x_q.size(1), -1)
376 t_barrel = torch.arange(t0, t1, device=k_star.device)
377 t_barrel = t_barrel[None, :].expand(k_star.size(0), t1 - t0)
379 torch.arange(k_star.size(0), device=k_star.device)[:, None] + t_barrel
381 k_star = k_star[l_barrel, t_barrel]
383 ######################################################################
384 # Compute the recurrent state
386 qw = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_qw)
388 v = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_v)
389 k = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_k)
395 ) / math.sqrt(self.w_qw.size(1))
397 aw = aw.softmax(dim=2) # nhlt
400 # We want all the memory lines to be used similarly
401 self.acc_attention += aw.sum(dim=(0, 1, 3)) # Sum accross NxHx_xT
402 self.acc_nb += aw.size(0) * aw.size(1) * aw.size(3)
404 aw = F.dropout(aw, self.attention_dropout, self.training)
406 A = 1 - aw.sum(dim=1) # nlt
408 V = torch.einsum("nhlt,nhtd->nltd", aw, v).contiguous()
409 K = torch.einsum("nhlt,nhtd->nltd", aw, k).contiguous()
415 V0 = self.rec_v[:, :, t0 - 1]
416 K0 = self.rec_k[:, :, t0 - 1]
418 self.rec_v[:, :, t0:t1] = pscan_shape(A, V, V0)
419 self.rec_k[:, :, t0:t1] = pscan_shape(A, K, K0)
421 ######################################################################
422 # compute the readout
424 qr = torch.einsum("ntc,hdc->nhtd", x_q[:, t0:t1], self.w_qr)
429 self.rec_k[:, :, t0:t1],
430 ) / math.sqrt(self.w_qr.size(1))
432 ar = ar.softmax(dim=2) # nhlt
434 ar = F.dropout(ar, self.attention_dropout, self.training)
439 self.rec_v[:, :, t0:t1],
442 self.cache_y[:, t0:t1] = y @ self.w_o
444 return BracketedSequence(self.cache_y, t0, t1 - t0, bs.init_cache)
447 ##############################
450 # Returns a tensor with an additional index at rank win_dim, that move
451 # along the same dimension as dim, on a domain {0...win_size-1}, and
452 # dim is restricted on a domain reduced by win_size-1 values.
455 def moving_window(x, dim, win_dim, win_size):
456 size, stride = x.size(), x.stride()
457 size = size[:dim] + (size[dim] - win_size + 1,) + size[dim + 1 :]
458 size = size[:win_dim] + (win_size,) + size[win_dim:]
459 stride = stride[:win_dim] + (stride[dim],) + stride[win_dim:]
461 return x.as_strided(size=size, stride=stride)
464 ##############################
467 class Caterpillar(nn.Module):
476 attention_dropout=0.0,
483 warnings.warn("Caterpillar", RuntimeWarning)
485 def randw(*d, amplitude=None):
486 if amplitude is None:
487 amplitude = 1 / math.sqrt(d[-1])
488 return nn.Parameter(amplitude * torch.randn(*d))
490 self.caterpillar_length = caterpillar_length
491 self.caterpillar_height = caterpillar_height
492 self.attention_dropout = attention_dropout
494 self.proba_gate_dropout = 0.0
496 default_b_G = kwargs.get("default_b_G")
497 if default_b_G is None:
498 default_b_G = -math.log(caterpillar_height - 1)
500 logger(f"default_b_G {default_b_G}")
502 self.w_G = randw(nb_heads, caterpillar_height, dim_model)
503 self.b_G = nn.Parameter(torch.full((nb_heads, caterpillar_height), default_b_G))
505 self.w_K = randw(nb_heads, dim_qk, dim_model)
506 self.w_V = randw(nb_heads, dim_v, dim_model)
507 self.w_Q = randw(nb_heads, dim_qk, dim_model)
508 self.w_O = randw(dim_v * nb_heads, dim_model)
510 self.init_K_rec = randw(
515 self.init_V_rec = randw(
521 def reset_inner_loss(self):
522 self.acc_attention = 0
525 def get_inner_loss(self):
526 # warnings.warn("l2 regularization", RuntimeWarning)
527 # return (self.acc_attention / self.acc_nb).pow(2).sum()
528 return torch.tensor([0], device=self.w_Q.device)
530 def forward(self, bs):
531 # Dimensions to make the source a bit clearer, that's needed
533 X, t0, t1 = bs.slice(), bs.first, bs.first + bs.nb
538 DV = self.w_V.size(1)
539 DK = self.w_K.size(1)
540 DM = self.w_O.size(1)
541 R = self.caterpillar_height
542 L = self.caterpillar_length
545 t0 >= L and (t1 - t0) % L == 0
546 ), f"bs.first should be greater than caterpillar_length, and bs.nb should be a multiple of caterpillar_length"
548 # We cache values to deal efficiently with auto-regression
551 self.rec_V = X.new_zeros(N, R, T, DV)
552 self.rec_K = X.new_zeros(N, R, T, DK)
553 # We start the recurrent sequences with optimizable
554 # initial values. No idea if it helps.
555 self.rec_V[:, :, t0 - L : t0] = self.init_V_rec[None, :, :, :]
556 self.rec_K[:, :, t0 - L : t0] = self.init_K_rec[None, :, :, :]
558 self.cache_Y = X.new_zeros(N, T, DM)
560 V = torch.einsum("ntc,hdc->nhtd", X, self.w_V)
561 K = torch.einsum("ntc,hdc->nhtd", X, self.w_K)
563 ######################################################################
564 # Compute the recurrent state
566 # This is the Gating sequence that modulates the storing of
567 # the new key and value in the R pairs of the current
568 # stack. There are R independent gating values, which means
569 # that the current K/V may be stored in multiple pairs of the
570 # recurrent state, or not at all.
573 torch.einsum("ntc,hrc->nhrt", X, self.w_G) + self.b_G[None, :, :, None]
576 # Clip the gating to avoid values greater than 1 when several
577 # heads hit the same row
579 G = G / G.sum(1, keepdim=True).clamp(min=1)
581 ######################################################################
582 # Roll the gating indexes
584 # warnings.warn("rotating barrel", RuntimeWarning)
586 # r_barrel = torch.arange(R, device=G.device)[None, None, :, None]
587 # t_barrel = torch.arange(t1 - t0, device=G.device)[None, None, None, :]
588 # r_barrel = (r_barrel + (t_barrel + t0) // L) % R
589 # G = G.gather(dim=2, index=r_barrel.expand_as(G))
591 ######################################################################
594 if self.training and self.proba_gate_dropout > 0.0:
595 # This is a better implementation of "flashbacks".
597 # G is NxHxExT where e is the caterpillar's row.
599 warnings.warn("gate dropout", RuntimeWarning)
603 (torch.rand(N, H, 1, t1 - t0, device=G.device).sort(dim=3).indices == 0)
608 dropout_tail = dropout_head.cumsum(dim=3) - dropout_head
611 torch.rand(N, 1, 1, 1, device=G.device) < self.proba_gate_dropout
614 dropout_head *= dropout_active
615 dropout_tail *= dropout_active
619 + dropout_head * (1 - epsilon - G.detach())
620 - dropout_tail * G.detach()
623 ######################################################################
625 # We prepare the arguments for the parallel scan
629 # warnings.warn("harmonic recurrence", RuntimeWarning)
630 # har = torch.arange(t0, t1, device = G.device).float() + 1
631 # A = har / (har + 1)
634 gated_V = torch.einsum("nhrt,nhtd->nrtd", G, V)
635 gated_K = torch.einsum("nhrt,nhtd->nrtd", G, K)
637 # We start from cached values, which matters in inference
639 init_rec_V = self.rec_V[:, :, t0 - L : t0]
640 init_rec_K = self.rec_K[:, :, t0 - L : t0]
642 #################################################################
645 # Here there is a trick: Since the stack at position t is
646 # computed by updating that at position t-L, the parallel
647 # scan operates with a period of L. To do so we split the
648 # sequence indexing in two axes, the second of size L, and
649 # run the parallel scan using the first as the sequence index.
651 A = A.unflatten(2, (-1, L))
652 gated_V = gated_V.unflatten(2, (-1, L))
653 gated_K = gated_K.unflatten(2, (-1, L))
655 next_V = pscan_dim(A, gated_V, init_rec_V, dim=2)
656 next_K = pscan_dim(A, gated_K, init_rec_K, dim=2)
658 self.rec_V[:, :, t0:t1] = next_V.flatten(2, 3)
659 self.rec_K[:, :, t0:t1] = next_K.flatten(2, 3)
661 ######################################################################
662 # compute the readout
664 Q = torch.einsum("ntc,hdc->nhtd", X, self.w_Q)
666 # We build tensors NxHxTxFxL where N is the sample index, H
667 # the head, T the time, F the row in the caterpillar, and L
668 # the column in the caterpillar
670 windowed_V = moving_window(
671 self.rec_V[:, :, t0 - L + 1 : t1], dim=2, win_dim=3, win_size=L
674 windowed_K = moving_window(
675 self.rec_K[:, :, t0 - L + 1 : t1], dim=2, win_dim=3, win_size=L
678 # We have an attention score for each of the RxL values
686 # softmax can operate only on one dimension, hence the
689 ar = ar.flatten(3).softmax(dim=3).view(ar.size())
691 ar = F.dropout(ar, self.attention_dropout, self.training)
693 # Compute the output for each head, flatten to concatenate
701 # Compute the final output
703 self.cache_Y[:, t0:t1] = Y @ self.w_O
705 return BracketedSequence(self.cache_Y, t0, t1 - t0, bs.init_cache)
708 ##############################
711 class QKVAttention(nn.Module):
719 attention_dropout=0.0,
726 return nn.Parameter(torch.randn(*d) / math.sqrt(d[-1]))
729 self.attention_dropout = attention_dropout
730 self.record_attention = False
732 self.w_q = randw(nb_heads, dim_qk, dim_model)
733 self.w_k = randw(nb_heads, dim_qk, dim_model)
734 self.w_v = randw(nb_heads, dim_v, dim_model)
735 self.w_o = randw(dim_v * nb_heads, dim_model)
737 def forward(self, bs):
741 self.causal or bs.complete()
742 ), "Partial evaluation is only possible for causal models"
745 self.cache_k = x_q.new_zeros(
746 x_q.size(0), self.w_k.size(0), x_q.size(1), self.w_k.size(1)
748 self.cache_v = x_q.new_zeros(
749 x_q.size(0), self.w_v.size(0), x_q.size(1), self.w_v.size(1)
751 self.cache_y = x_q.new_zeros(x_q.size(0), x_q.size(1), self.w_o.size(1))
753 q = torch.einsum("ntc,hdc->nhtd", x_q[:, bs.first : bs.first + bs.nb], self.w_q)
755 self.cache_k[:, :, bs.first : bs.first + bs.nb] = torch.einsum(
756 "ntc,hdc->nhtd", x_q[:, bs.first : bs.first + bs.nb], self.w_k
758 self.cache_v[:, :, bs.first : bs.first + bs.nb] = torch.einsum(
759 "ntc,hdc->nhtd", x_q[:, bs.first : bs.first + bs.nb], self.w_v
763 "nhtd,nhsd->nhts", q, self.cache_k[:, :, : bs.first + bs.nb]
764 ) / math.sqrt(self.w_q.size(1))
768 self.cache_attzero = (
769 torch.arange(x_q.size(1), device=q.device)[None, None, :, None]
770 < torch.arange(x_q.size(1), device=q.device)[None, None, None, :]
774 :, :, bs.first : bs.first + bs.nb, : bs.first + bs.nb
781 if self.record_attention:
784 a = F.dropout(a, self.attention_dropout, self.training)
787 "nhts,nhsd->nthd", a, self.cache_v[:, :, : bs.first + bs.nb]
790 self.cache_y[:, bs.first : bs.first + bs.nb] = y @ self.w_o
792 return BracketedSequence(self.cache_y, bs.first, bs.nb, bs.init_cache)
795 ##############################
798 class MyGPT(nn.Module):
808 caterpillar_height=None,
812 attention_layer="kvrec",
818 assert attention_layer in {
823 }, f"Unknown attention operator {attention_layer}."
825 if attention_layer == "caterpillar":
826 assert nb_lines % caterpillar_height == 0
827 self.caterpillar_length = nb_lines // caterpillar_height
828 self.caterpillar_height = caterpillar_height
830 self.caterpillar_length = -1
831 self.caterpillar_height = -1
833 assert dim_model % nb_heads == 0
835 self.embedding = nn.Sequential(
836 CacheWrapper(nn.Embedding(vocabulary_size, dim_model), nn.Dropout(dropout)),
837 AddPositionalEncoding(len_max),
843 if attention_layer == "mha":
847 dim_v=dim_model // nb_heads,
850 attention_dropout=dropout,
854 elif attention_layer == "dumbrec":
858 dim_v=dim_model // nb_heads,
861 attention_dropout=dropout,
865 elif attention_layer == "kvrec":
869 dim_v=dim_model // nb_heads,
872 attention_dropout=dropout,
876 elif attention_layer == "caterpillar":
880 dim_v=dim_model // nb_heads,
882 caterpillar_length=self.caterpillar_length,
883 caterpillar_height=self.caterpillar_height,
884 attention_dropout=dropout,
889 raise ValueError(f"Unknown attention type {attention_layer}.")
891 for b in range(nb_blocks):
894 CacheWrapper(nn.LayerNorm((dim_model,))),
899 nn.LayerNorm((dim_model,)),
900 nn.Linear(in_features=dim_model, out_features=dim_hidden),
902 nn.Linear(in_features=dim_hidden, out_features=dim_model),
908 self.trunk = nn.Sequential(*trunk_blocks)
910 self.readout = CacheWrapper(
911 nn.Linear(in_features=dim_model, out_features=vocabulary_size)
914 with torch.no_grad():
915 for m in self.modules():
916 if isinstance(m, nn.Embedding):
917 m.weight.normal_(mean=0, std=2e-2)
918 elif isinstance(m, nn.LayerNorm):
922 self.reset_inner_loss()
924 def forward(self, bs):
925 bs = BracketedSequence(F.pad(bs.x, (1, -1)), bs.first, bs.nb, bs.init_cache)
927 # To make the code simpler in the Caterpillar layer, we pad
928 # here. It's unclear if/how much it hurts computationaly by
929 # increasing the sequence length for the other layers
931 if self.caterpillar_length > 0:
933 if bs.nb % self.caterpillar_length > 0:
934 bs.nb += self.caterpillar_length - bs.nb % self.caterpillar_length
936 bs = BracketedSequence(
937 F.pad(bs.x, (self.caterpillar_length, self.caterpillar_length)),
938 bs.first + self.caterpillar_length,
943 bs = self.embedding(bs)
945 bs = self.readout(bs)
947 if self.caterpillar_length > 0:
948 bs = BracketedSequence(
949 F.pad(bs.x, (0, 0, -self.caterpillar_length, -self.caterpillar_length)),
950 bs.first - self.caterpillar_length,
957 # ar_mask is a tensor with 0s and 1s, of same shape as input, with
958 # 1s where tokens should be generated. The others are kept
961 def masked_inplace_autoregression(
965 forbidden_tokens=None,
966 deterministic_synthesis=False,
968 input = input_src.to(self.readout.f.weight.device)
969 ar_mask = ar_mask_src.to(self.readout.f.weight.device)
970 to_generate = (ar_mask.sum(0) > 0).nonzero()
971 if to_generate.min() > 0:
973 BracketedSequence(input, 0, to_generate.min(), True)
974 ) # Needed to initialize the model's cache
975 for s in range(to_generate.min(), to_generate.max() + 1):
976 output = self(BracketedSequence(input, s, 1, s == 0)).x
977 logits = output[:, s]
978 if forbidden_tokens is not None:
979 logits = logits.masked_fill(forbidden_tokens, float("-inf"))
980 if deterministic_synthesis:
981 t_next = logits.argmax(1)
983 dist = torch.distributions.categorical.Categorical(logits=logits)
984 t_next = dist.sample()
985 input[:, s] = ar_mask[:, s] * t_next + (1 - ar_mask[:, s]) * input[:, s]
987 input_src.copy_(input)
989 def reset_inner_loss(self):
990 for m in self.modules():
991 if m is not self and hasattr(m, "reset_inner_loss"):
994 def get_inner_loss(self):
995 l = torch.tensor([0.0], device=self.readout.f.weight.device)
996 for m in self.modules():
997 if m is not self and hasattr(m, "get_inner_loss"):
998 l += m.get_inner_loss()
1001 def record_attention(self, v=True):
1002 for m in self.modules():
1003 if isinstance(m, QKVAttention):
1004 m.record_attention = v
1006 def retrieve_attention(self):
1008 for m in self.modules():
1009 if isinstance(m, QKVAttention):
1014 ######################################################################
1016 if __name__ == "__main__":
1017 print("Basic check.")
1024 caterpillar_length=7,
1025 caterpillar_height=3,
1026 attention_dropout=0.0,
1029 m.reset_inner_loss()
1030 x = torch.randn(1, 21 + 2 * 7, 4)
1031 y1 = m(BracketedSequence(x, first=7, nb=21, init_cache=True)).x[:, 7:28]
1032 y2 = m(BracketedSequence(x, first=7, nb=21, init_cache=True)).x[:, 7:28]
1033 y3a = m(BracketedSequence(x, first=7, nb=14, init_cache=True)).x[:, 7:21]
1034 y3b = m(BracketedSequence(x, first=21, nb=7, init_cache=False)).x[:, 21:28]
1035 print((y1 - y2).abs().max())
1036 print((y1 - torch.cat([y3a, y3b], dim=1)).abs().max())
1039 device = torch.device("cuda" if torch.cuda.is_available() else "cpu")
1041 vocabulary_size = 128
1042 x = torch.randint(vocabulary_size, (6, 1024))
1045 vocabulary_size=vocabulary_size,
1061 # import torchvision.models as models
1062 # from torch.profiler import profile, record_function, ProfilerActivity
1064 # with profile(activities=[ProfilerActivity.CPU, ProfilerActivity.CUDA], profile_memory=True, record_shapes=True) as prof:
1065 # with record_function("model_inference"):
1069 start_time = time.perf_counter()
1071 model(BracketedSequence(x))
1072 duration = time.perf_counter() - start_time
1076 # print(prof.key_averages().table(sort_by="cpu_time_total", row_limit=10))
1077 # print(prof.key_averages().table(sort_by="cuda_time_total", row_limit=10))
1079 # print("##############################################################")
1080 # y2 = torch.randn_like(y1)
1081 # for s in range(x.size(1)):
1082 # z = model(BracketedSequence(x, s, 1))
1083 # y2[:, s : s + 1] = z.slice()
1085 # print(f"error={((y1 - y2).norm() / (y1.norm() + y2.norm())).item()}")
1087 ######################################################################