block.py 17 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397
  1. # Copyright (c) 2024, Tri Dao.
  2. from functools import partial
  3. from typing import Optional
  4. import torch
  5. import torch.nn as nn
  6. import torch.nn.functional as F
  7. from torch import Tensor
  8. from torchvision.ops import StochasticDepth
  9. from flash_attn.modules.mha import MHA
  10. from flash_attn.modules.mlp import Mlp
  11. try:
  12. from flash_attn.ops.triton.layer_norm import layer_norm_fn, RMSNorm
  13. except ImportError:
  14. layer_norm_fn, RMSNorm = None, None
  15. class Block(nn.Module):
  16. def __init__(
  17. self,
  18. dim,
  19. mixer_cls=None,
  20. mlp_cls=None,
  21. norm_cls=nn.LayerNorm,
  22. dropout_cls=nn.Dropout,
  23. prenorm=True,
  24. resid_dropout1=0.0,
  25. resid_dropout2=0.0,
  26. drop_path1=0.0,
  27. drop_path2=0.0,
  28. fused_dropout_add_ln=False,
  29. return_residual=False,
  30. residual_in_fp32=False,
  31. sequence_parallel=False,
  32. mark_shared_params=False,
  33. ):
  34. """
  35. For prenorm=True, this Block has a slightly different structure compared to a regular
  36. prenorm Transformer block.
  37. The standard block is: LN -> MHA -> Dropout -> Add -> LN -> MLP -> Dropout -> Add.
  38. [Ref: https://arxiv.org/abs/2002.04745]
  39. Here we have: Dropout -> Add -> LN -> MHA -> Dropout -> Add -> LN -> MLP, returning both
  40. the hidden_states (output of the MLP) and the residual.
  41. This is for performance reasons, as we can fuse the dropout, add and LayerNorm.
  42. The residual needs to be provided (except for the very first block).
  43. For prenorm=False, this Block has the same structure as a regular postnorm Transformer
  44. block: MHA -> Dropout -> Add -> LN -> MLP -> Dropout -> Add -> LN.
  45. return_residual: whether each of the sub-layers (mixer and mlp) will return the residual.
  46. This is for performance reason: for post-norm architecture, returning the input allows us
  47. to fuse the backward of nn.Linear with the residual connection.
  48. """
  49. super().__init__()
  50. self.prenorm = prenorm
  51. self.fused_dropout_add_ln = fused_dropout_add_ln
  52. self.return_residual = return_residual
  53. self.residual_in_fp32 = residual_in_fp32
  54. if self.residual_in_fp32:
  55. assert self.prenorm, "residual_in_fp32 is only compatible with prenorm=True"
  56. if mixer_cls is None:
  57. mixer_cls = partial(MHA, num_heads=dim // 64)
  58. if mlp_cls is None:
  59. mlp_cls = partial(Mlp, hidden_features=4 * dim)
  60. self.mixer = mixer_cls(dim)
  61. self.dropout1 = dropout_cls(resid_dropout1)
  62. self.drop_path1 = StochasticDepth(drop_path1, mode="row")
  63. self.norm1 = norm_cls(dim)
  64. self.mlp = mlp_cls(dim)
  65. if not isinstance(self.mlp, nn.Identity):
  66. self.dropout2 = dropout_cls(resid_dropout2)
  67. self.drop_path2 = StochasticDepth(drop_path2, mode="row")
  68. self.norm2 = norm_cls(dim)
  69. if self.fused_dropout_add_ln:
  70. assert layer_norm_fn is not None, "Triton is not installed"
  71. assert isinstance(self.norm1, (nn.LayerNorm, RMSNorm)) and isinstance(
  72. self.dropout1, nn.Dropout
  73. )
  74. # TD [2023-01-07]: TODO: During training, if sequence_parallel is False and dropout != 0.0,
  75. # then the input to each worker in the tensor parallel group will be different.
  76. # This would produce wrong outputs? Somehow we'd need to sync the RNG state across workers.
  77. # For now this is not an issue because we always use sequence_parallel=True during training
  78. # and only use sequence_parallel=False during inference.
  79. # Mark the norm parameters as "sequence_parallel" so that we run all-reduce on their grads.
  80. if sequence_parallel:
  81. for p in self.norm1.parameters():
  82. p._sequence_parallel = True
  83. if hasattr(self, "norm2"):
  84. for p in self.norm2.parameters():
  85. p._sequence_parallel = True
  86. # Mark the norm parameters as "shared_params" so that we sync their values at init.
  87. if mark_shared_params:
  88. for p in self.norm1.parameters():
  89. p._shared_params = True
  90. if hasattr(self, "norm2"):
  91. for p in self.norm2.parameters():
  92. p._shared_params = True
  93. def allocate_inference_cache(self, batch_size, max_seqlen, dtype=None, **kwargs):
  94. return self.mixer.allocate_inference_cache(batch_size, max_seqlen, dtype=dtype, **kwargs)
  95. def forward(
  96. self,
  97. hidden_states: Tensor,
  98. residual: Optional[Tensor] = None,
  99. mixer_subset=None,
  100. mixer_kwargs=None,
  101. ):
  102. r"""Pass the input through the encoder layer.
  103. Args:
  104. hidden_states: the sequence to the encoder layer (required).
  105. residual: if postnorm, residual=None, If prenorm, hidden_states = Attn/MLP(LN(residual))
  106. mixer_subset: for cross-attention only. If not None, will take a subset of x
  107. before applying the query projection. Useful for e.g., ViT where we only care
  108. about the CLS token in the last layer.
  109. """
  110. if self.prenorm:
  111. if not self.fused_dropout_add_ln:
  112. dropped = self.drop_path1(self.dropout1(hidden_states))
  113. residual = (dropped + residual) if residual is not None else dropped
  114. hidden_states = self.norm1(residual.to(dtype=self.norm1.weight.dtype))
  115. if self.residual_in_fp32:
  116. residual = residual.to(torch.float32)
  117. else:
  118. if self.drop_path1.p == 0 or not self.training:
  119. rowscale1 = None
  120. else:
  121. rowscale1 = self.drop_path1(
  122. torch.ones(
  123. hidden_states.shape[:-1],
  124. device=hidden_states.device,
  125. dtype=hidden_states.dtype,
  126. )
  127. )
  128. hidden_states, residual = layer_norm_fn(
  129. hidden_states,
  130. self.norm1.weight,
  131. self.norm1.bias,
  132. residual=residual,
  133. eps=self.norm1.eps,
  134. dropout_p=self.dropout1.p if self.training else 0.0,
  135. rowscale=rowscale1,
  136. prenorm=True,
  137. residual_in_fp32=self.residual_in_fp32,
  138. is_rms_norm=isinstance(self.norm1, RMSNorm)
  139. )
  140. if mixer_kwargs is None:
  141. mixer_kwargs = {}
  142. if mixer_subset is not None:
  143. mixer_kwargs["mixer_subset"] = mixer_subset
  144. hidden_states = self.mixer(hidden_states, **mixer_kwargs)
  145. if mixer_subset is not None:
  146. residual = residual[:, mixer_subset]
  147. if not isinstance(self.mlp, nn.Identity):
  148. if not self.fused_dropout_add_ln:
  149. dropped = self.drop_path2(self.dropout2(hidden_states))
  150. residual = (dropped + residual) if residual is not None else dropped
  151. hidden_states = self.norm2(residual.to(dtype=self.norm2.weight.dtype))
  152. if self.residual_in_fp32:
  153. residual = residual.to(torch.float32)
  154. else:
  155. if self.drop_path2.p == 0 or not self.training:
  156. rowscale2 = None
  157. else:
  158. rowscale2 = self.drop_path2(
  159. torch.ones(
  160. hidden_states.shape[:-1],
  161. device=hidden_states.device,
  162. dtype=hidden_states.dtype,
  163. )
  164. )
  165. hidden_states, residual = layer_norm_fn(
  166. hidden_states,
  167. self.norm2.weight,
  168. self.norm2.bias,
  169. residual=residual,
  170. eps=self.norm2.eps,
  171. dropout_p=self.dropout2.p if self.training else 0.0,
  172. rowscale=rowscale2,
  173. prenorm=True,
  174. residual_in_fp32=self.residual_in_fp32,
  175. is_rms_norm=isinstance(self.norm2, RMSNorm)
  176. )
  177. hidden_states = self.mlp(hidden_states)
  178. return hidden_states, residual
  179. else:
  180. assert residual is None
  181. mixer_out = self.mixer(
  182. hidden_states, **(mixer_kwargs if mixer_kwargs is not None else {})
  183. )
  184. if self.return_residual: # mixer out is actually a pair here
  185. mixer_out, hidden_states = mixer_out
  186. if not self.fused_dropout_add_ln:
  187. hidden_states = self.norm1(
  188. (self.drop_path1(self.dropout1(mixer_out)) + hidden_states).to(
  189. dtype=self.norm1.weight.dtype
  190. )
  191. )
  192. else:
  193. if self.drop_path1.p == 0 or not self.training:
  194. rowscale1 = None
  195. else:
  196. rowscale1 = self.drop_path1(
  197. torch.ones(
  198. mixer_out.shape[:-1], device=mixer_out.device, dtype=mixer_out.dtype
  199. )
  200. )
  201. hidden_states = layer_norm_fn(
  202. mixer_out,
  203. self.norm1.weight,
  204. self.norm1.bias,
  205. residual=hidden_states,
  206. eps=self.norm1.eps,
  207. dropout_p=self.dropout1.p if self.training else 0.0,
  208. rowscale=rowscale1,
  209. prenorm=False,
  210. is_rms_norm=isinstance(self.norm1, RMSNorm)
  211. )
  212. if not isinstance(self.mlp, nn.Identity):
  213. mlp_out = self.mlp(hidden_states)
  214. if self.return_residual: # mlp out is actually a pair here
  215. mlp_out, hidden_states = mlp_out
  216. if not self.fused_dropout_add_ln:
  217. hidden_states = self.norm2(
  218. (self.drop_path2(self.dropout2(mlp_out)) + hidden_states).to(
  219. dtype=self.norm2.weight.dtype
  220. )
  221. )
  222. else:
  223. if self.drop_path2.p == 0 or not self.training:
  224. rowscale2 = None
  225. else:
  226. rowscale2 = self.drop_path2(
  227. torch.ones(
  228. mlp_out.shape[:-1], device=mlp_out.device, dtype=mlp_out.dtype
  229. )
  230. )
  231. hidden_states = layer_norm_fn(
  232. mlp_out,
  233. self.norm2.weight,
  234. self.norm2.bias,
  235. residual=hidden_states,
  236. eps=self.norm2.eps,
  237. dropout_p=self.dropout2.p if self.training else 0.0,
  238. rowscale=rowscale2,
  239. prenorm=False,
  240. is_rms_norm=isinstance(self.norm2, RMSNorm)
  241. )
  242. return hidden_states
  243. class ParallelBlock(nn.Module):
  244. """The attention (mixer) and MLP blocks are done in parallel, similar to GPT-J, GPT-NeoX,
  245. and PaLM.
  246. """
  247. def __init__(
  248. self,
  249. dim,
  250. mixer_cls=None,
  251. mlp_cls=None,
  252. norm_cls=nn.LayerNorm,
  253. dropout_cls=nn.Dropout,
  254. resid_dropout1=0.0,
  255. resid_dropout2=0.0,
  256. tied_norm=False,
  257. fused_dropout_add_ln=False,
  258. residual_in_fp32=False,
  259. sequence_parallel=False,
  260. mark_shared_params=False,
  261. ):
  262. """
  263. This Block has a slightly different structure compared to a regular
  264. prenorm Transformer block.
  265. The standard block is: LN -> MHA / MLP -> Dropout -> Add.
  266. [Ref: https://arxiv.org/abs/2002.04745]
  267. Here we have: Dropout -> Add -> LN -> MHA / MLP, returning both
  268. the hidden_states (output1 of the MHA / MLP) and the residual.
  269. This is for performance reasons, as we can fuse the dropout, add and LayerNorm.
  270. The residual needs to be provided (except for the very first block).
  271. """
  272. super().__init__()
  273. self.tied_norm = tied_norm
  274. self.fused_dropout_add_ln = fused_dropout_add_ln
  275. self.residual_in_fp32 = residual_in_fp32
  276. if mixer_cls is None:
  277. mixer_cls = partial(MHA, num_heads=dim // 64)
  278. if mlp_cls is None:
  279. mlp_cls = partial(Mlp, hidden_features=4 * dim)
  280. self.mixer = mixer_cls(dim)
  281. self.dropout1 = dropout_cls(resid_dropout1)
  282. self.norm1 = norm_cls(dim)
  283. self.mlp = mlp_cls(dim)
  284. self.dropout2 = dropout_cls(resid_dropout2)
  285. if not self.tied_norm:
  286. self.norm2 = norm_cls(dim)
  287. if self.fused_dropout_add_ln:
  288. assert layer_norm_fn is not None, "Triton is not installed"
  289. assert isinstance(self.norm1, (nn.LayerNorm, RMSNorm)) and isinstance(
  290. self.dropout1, nn.Dropout
  291. )
  292. # TD [2023-01-07]: TODO: During training, if sequence_parallel is False and dropout != 0.0,
  293. # then the input to each worker in the tensor parallel group will be different.
  294. # This would produce wrong outputs? Somehow we'd need to sync the RNG state across workers.
  295. # For now this is not an issue because we always use sequence_parallel=True during training
  296. # and only use sequence_parallel=False during inference.
  297. # Mark the norm parameters as "sequence_parallel" so that we run all-reduce on their grads.
  298. if sequence_parallel:
  299. for p in self.norm1.parameters():
  300. p._sequence_parallel = True
  301. if hasattr(self, "norm2"):
  302. for p in self.norm2.parameters():
  303. p._sequence_parallel = True
  304. # Mark the norm parameters as "shared_params" so that we sync their values at init.
  305. if mark_shared_params:
  306. for p in self.norm1.parameters():
  307. p._shared_params = True
  308. if hasattr(self, "norm2"):
  309. for p in self.norm2.parameters():
  310. p._shared_params = True
  311. def allocate_inference_cache(self, batch_size, max_seqlen, dtype=None, **kwargs):
  312. return self.mixer.allocate_inference_cache(batch_size, max_seqlen, dtype=dtype, **kwargs)
  313. def forward(
  314. self,
  315. hidden_states1: Tensor,
  316. hidden_states2: Optional[Tensor] = None,
  317. residual: Optional[Tensor] = None,
  318. mixer_kwargs=None,
  319. ):
  320. r"""Pass the input through the encoder layer.
  321. Args:
  322. hidden_states1: the output of the previous attention (mixer) or embedding layer.
  323. hidden_states2: the output of the previous MLP layer (if None, will use hidden_states1).
  324. residual.
  325. """
  326. # TODO: Ideally we should only do the allgather / allreduce once for
  327. # the Linear to MLP & Attention
  328. if not self.fused_dropout_add_ln:
  329. dropped1 = self.dropout1(hidden_states1)
  330. # For the very 1st block, we only want 1 dropout, not two different dropouts
  331. if hidden_states2 is not None:
  332. dropped2 = self.dropout2(hidden_states2)
  333. residual = (
  334. (residual + dropped1 + dropped2)
  335. if residual is not None
  336. else dropped1 + dropped2
  337. )
  338. else:
  339. residual = (residual + dropped1) if residual is not None else dropped1
  340. hidden_states1 = self.norm1(residual.to(dtype=self.norm1.weight.dtype))
  341. hidden_states2 = (
  342. self.norm2(residual.to(dtype=self.norm2.weight.dtype))
  343. if not self.tied_norm
  344. else hidden_states1
  345. )
  346. if self.residual_in_fp32:
  347. residual = residual.to(torch.float32)
  348. else:
  349. weight2, bias2 = (
  350. (self.norm2.weight, self.norm2.bias) if not self.tied_norm else (None, None)
  351. )
  352. hidden_states1, *rest, residual = layer_norm_fn(
  353. hidden_states1,
  354. self.norm1.weight,
  355. self.norm1.bias,
  356. residual=residual,
  357. x1=hidden_states2,
  358. weight1=weight2,
  359. bias1=bias2,
  360. eps=self.norm1.eps,
  361. dropout_p=self.dropout1.p if self.training else 0.0,
  362. prenorm=True,
  363. residual_in_fp32=self.residual_in_fp32,
  364. is_rms_norm=isinstance(self.norm1, RMSNorm)
  365. )
  366. if self.tied_norm:
  367. hidden_states2 = hidden_states1
  368. else:
  369. hidden_states2, = rest
  370. if mixer_kwargs is None:
  371. mixer_kwargs = {}
  372. hidden_states1 = self.mixer(hidden_states1, **mixer_kwargs)
  373. hidden_states2 = self.mlp(hidden_states2)
  374. return hidden_states1, hidden_states2, residual