cache_utils.py 20 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434
  1. from dataclasses import dataclass
  2. from typing import Any, Dict, List, Optional, Tuple
  3. import torch
  4. from transformers.configuration_utils import PretrainedConfig
  5. from transformers.utils import logging
  6. logger = logging.get_logger(__name__)
  7. @dataclass
  8. class Cache:
  9. """
  10. Base, abstract class for all caches. The actual data structure is specific to each subclass.
  11. """
  12. def update(
  13. self,
  14. key_states: torch.Tensor,
  15. value_states: torch.Tensor,
  16. layer_idx: int,
  17. cache_kwargs: Optional[Dict[str, Any]] = None,
  18. ) -> Tuple[torch.Tensor, torch.Tensor]:
  19. """
  20. Updates the cache with the new `key_states` and `value_states` for the layer `layer_idx`.
  21. Parameters:
  22. key_states (`torch.Tensor`):
  23. The new key states to cache.
  24. value_states (`torch.Tensor`):
  25. The new value states to cache.
  26. layer_idx (`int`):
  27. The index of the layer to cache the states for.
  28. cache_kwargs (`Dict[str, Any]`, `optional`):
  29. Additional arguments for the cache subclass. These are specific to each subclass and allow new types of
  30. cache to be created.
  31. Return:
  32. A tuple containing the updated key and value states.
  33. """
  34. raise NotImplementedError("Make sure to implement `update` in a subclass.")
  35. def get_seq_length(self, layer_idx: Optional[int] = 0) -> int:
  36. """Returns the sequence length of the cached states. A layer index can be optionally passed."""
  37. raise NotImplementedError("Make sure to implement `get_seq_length` in a subclass.")
  38. def get_max_length(self) -> Optional[int]:
  39. """Returns the maximum sequence length of the cached states, if there is any."""
  40. raise NotImplementedError("Make sure to implement `get_max_length` in a subclass.")
  41. def get_usable_length(self, new_seq_length: int, layer_idx: Optional[int] = 0) -> int:
  42. """Given the sequence length of the new inputs, returns the usable length of the cache."""
  43. # Cache without size limit -> all cache is usable
  44. # Cache with size limit -> if the length cache plus the length of the new inputs is larger the maximum cache
  45. # length, we will need to evict part of the cache (and thus not all cache is usable)
  46. max_length = self.get_max_length()
  47. previous_seq_length = self.get_seq_length(layer_idx)
  48. if max_length is not None and previous_seq_length + new_seq_length > max_length:
  49. return max_length - new_seq_length
  50. return previous_seq_length
  51. @property
  52. def seen_tokens(self):
  53. logger.warning_once(
  54. "The `seen_tokens` attribute is deprecated and will be removed in v4.41. Use the `cache_position` "
  55. "model input instead."
  56. )
  57. if hasattr(self, "_seen_tokens"):
  58. return self._seen_tokens
  59. else:
  60. return None
  61. class DynamicCache(Cache):
  62. """
  63. A cache that grows dynamically as more tokens are generated. This is the default for generative models.
  64. It stores the Key and Value states as a list of tensors, one for each layer. The expected shape for each tensor is
  65. `[batch_size, num_heads, seq_len, head_dim]`.
  66. """
  67. def __init__(self) -> None:
  68. self.key_cache: List[torch.Tensor] = []
  69. self.value_cache: List[torch.Tensor] = []
  70. self._seen_tokens = 0 # Used in `generate` to keep tally of how many tokens the cache has seen
  71. def __getitem__(self, layer_idx: int) -> List[Tuple[torch.Tensor]]:
  72. """
  73. Support for backwards-compatible `past_key_value` indexing, e.g. `past_key_value[0][0].shape[2]` to get the
  74. sequence length.
  75. """
  76. if layer_idx < len(self):
  77. return (self.key_cache[layer_idx], self.value_cache[layer_idx])
  78. else:
  79. raise KeyError(f"Cache only has {len(self)} layers, attempted to access layer with index {layer_idx}")
  80. def __iter__(self):
  81. """
  82. Support for backwards-compatible `past_key_value` iteration, e.g. `for x in past_key_value:` to iterate over
  83. keys and values
  84. """
  85. for layer_idx in range(len(self)):
  86. yield (self.key_cache[layer_idx], self.value_cache[layer_idx])
  87. def __len__(self):
  88. """
  89. Support for backwards-compatible `past_key_value` length, e.g. `len(past_key_value)`. This value corresponds
  90. to the number of layers in the model.
  91. """
  92. return len(self.key_cache)
  93. def update(
  94. self,
  95. key_states: torch.Tensor,
  96. value_states: torch.Tensor,
  97. layer_idx: int,
  98. cache_kwargs: Optional[Dict[str, Any]] = None,
  99. ) -> Tuple[torch.Tensor, torch.Tensor]:
  100. """
  101. Updates the cache with the new `key_states` and `value_states` for the layer `layer_idx`.
  102. Parameters:
  103. key_states (`torch.Tensor`):
  104. The new key states to cache.
  105. value_states (`torch.Tensor`):
  106. The new value states to cache.
  107. layer_idx (`int`):
  108. The index of the layer to cache the states for.
  109. cache_kwargs (`Dict[str, Any]`, `optional`):
  110. Additional arguments for the cache subclass. No additional arguments are used in `DynamicCache`.
  111. Return:
  112. A tuple containing the updated key and value states.
  113. """
  114. # Update the number of seen tokens
  115. if layer_idx == 0:
  116. self._seen_tokens += key_states.shape[-2]
  117. # Update the cache
  118. print(len(self.key_cache), layer_idx)
  119. if len(self.key_cache) <= layer_idx:
  120. self.key_cache.append(key_states)
  121. self.value_cache.append(value_states)
  122. else:
  123. self.key_cache[layer_idx] = torch.cat([self.key_cache[layer_idx], key_states], dim=-2)
  124. self.value_cache[layer_idx] = torch.cat([self.value_cache[layer_idx], value_states], dim=-2)
  125. return self.key_cache[layer_idx], self.value_cache[layer_idx]
  126. def get_seq_length(self, layer_idx: Optional[int] = 0) -> int:
  127. """Returns the sequence length of the cached states. A layer index can be optionally passed."""
  128. if len(self.key_cache) <= layer_idx:
  129. return 0
  130. return self.key_cache[layer_idx].shape[-2]
  131. def get_max_length(self) -> Optional[int]:
  132. """Returns the maximum sequence length of the cached states. DynamicCache does not have a maximum length."""
  133. return None
  134. def reorder_cache(self, beam_idx: torch.LongTensor):
  135. """Reorders the cache for beam search, given the selected beam indices."""
  136. for layer_idx in range(len(self.key_cache)):
  137. device = self.key_cache[layer_idx].device
  138. self.key_cache[layer_idx] = self.key_cache[layer_idx].index_select(0, beam_idx.to(device))
  139. device = self.value_cache[layer_idx].device
  140. self.value_cache[layer_idx] = self.value_cache[layer_idx].index_select(0, beam_idx.to(device))
  141. def to_legacy_cache(self) -> Tuple[Tuple[torch.Tensor], Tuple[torch.Tensor]]:
  142. """Converts the `DynamicCache` instance into the its equivalent in the legacy cache format."""
  143. legacy_cache = ()
  144. for layer_idx in range(len(self)):
  145. legacy_cache += ((self.key_cache[layer_idx], self.value_cache[layer_idx]),)
  146. return legacy_cache
  147. @classmethod
  148. def from_legacy_cache(cls, past_key_values: Optional[Tuple[Tuple[torch.FloatTensor]]] = None) -> "DynamicCache":
  149. """Converts a cache in the legacy cache format into an equivalent `DynamicCache`."""
  150. cache = cls()
  151. if past_key_values is not None:
  152. for layer_idx in range(len(past_key_values)):
  153. key_states, value_states = past_key_values[layer_idx]
  154. cache.update(key_states, value_states, layer_idx)
  155. return cache
  156. class SinkCache(Cache):
  157. """
  158. A cache that as described in the [Attention Sinks paper](https://arxiv.org/abs/2309.17453). It allows the model to
  159. generate beyond the length of its context window, without losing fluency in the conversation. As it discards past
  160. tokens, the model will lose the ability to generate tokens that depend on the context that was discarded.
  161. It stores the Key and Value states as a list of tensors, one for each layer. The expected shape for each tensor is
  162. `[batch_size, num_heads, seq_len, head_dim]`.
  163. Parameters:
  164. window_length (`int`):
  165. The length of the context window.
  166. num_sink_tokens (`int`):
  167. The number of sink tokens. See the original paper for more information.
  168. """
  169. def __init__(self, window_length: int, num_sink_tokens: int) -> None:
  170. self.key_cache: List[torch.Tensor] = []
  171. self.value_cache: List[torch.Tensor] = []
  172. self.window_length = window_length
  173. self.num_sink_tokens = num_sink_tokens
  174. self.cos_sin_cache = {}
  175. self._seen_tokens = 0 # Used in `generate` to keep tally of how many tokens the cache has seen
  176. @staticmethod
  177. def _rotate_half(x):
  178. x1 = x[..., : x.shape[-1] // 2]
  179. x2 = x[..., x.shape[-1] // 2 :]
  180. return torch.cat((-x2, x1), dim=-1)
  181. def _apply_key_rotary_pos_emb(
  182. self, key_states: torch.Tensor, cos: torch.Tensor, sin: torch.Tensor
  183. ) -> torch.Tensor:
  184. rotated_key_states = (key_states * cos) + (self._rotate_half(key_states) * sin)
  185. return rotated_key_states
  186. def _get_rerotation_cos_sin(
  187. self, key_states: torch.Tensor, cos: torch.Tensor, sin: torch.Tensor
  188. ) -> Tuple[torch.Tensor, torch.Tensor]:
  189. if key_states.shape[-2] not in self.cos_sin_cache:
  190. # Upcast to float32 temporarily for better accuracy
  191. cos = cos.to(torch.float32)
  192. sin = sin.to(torch.float32)
  193. # Compute the cos and sin required for back- and forward-rotating to one position earlier in the sequence
  194. original_cos = cos[self.num_sink_tokens + key_states.shape[-2] :]
  195. shifted_cos = cos[self.num_sink_tokens : -key_states.shape[-2]]
  196. original_sin = sin[self.num_sink_tokens + key_states.shape[-2] :]
  197. shifted_sin = sin[self.num_sink_tokens : -key_states.shape[-2]]
  198. rerotation_cos = original_cos * shifted_cos + original_sin * shifted_sin
  199. rerotation_sin = -original_sin * shifted_cos + original_cos * shifted_sin
  200. self.cos_sin_cache[key_states.shape[-2]] = (
  201. rerotation_cos.to(key_states.dtype).unsqueeze(0),
  202. rerotation_sin.to(key_states.dtype).unsqueeze(0),
  203. )
  204. return self.cos_sin_cache[key_states.shape[-2]]
  205. def get_seq_length(self, layer_idx: Optional[int] = 0) -> int:
  206. """Returns the sequence length of the cached states. A layer index can be optionally passed."""
  207. # Workaround to make 'key_states.shape[-2] + past_key_value.get_seq_length(self.layer_idx)' <= window_length
  208. if len(self.key_cache) <= layer_idx:
  209. return 0
  210. return self.key_cache[layer_idx].shape[-2]
  211. def get_max_length(self) -> Optional[int]:
  212. """Returns the maximum sequence length of the cached states."""
  213. return self.window_length
  214. def update(
  215. self,
  216. key_states: torch.Tensor,
  217. value_states: torch.Tensor,
  218. layer_idx: int,
  219. cache_kwargs: Optional[Dict[str, Any]] = None,
  220. ) -> Tuple[torch.Tensor, torch.Tensor]:
  221. """
  222. Updates the cache with the new `key_states` and `value_states` for the layer `layer_idx`.
  223. Parameters:
  224. key_states (`torch.Tensor`):
  225. The new key states to cache.
  226. value_states (`torch.Tensor`):
  227. The new value states to cache.
  228. layer_idx (`int`):
  229. The index of the layer to cache the states for.
  230. cache_kwargs (`Dict[str, Any]`, `optional`):
  231. Additional arguments for the cache subclass. The following arguments can be used in `SinkCache`: `sin`,
  232. `cos` and `partial_rotation_size`. These arguments are used with models using RoPE, to recompute the
  233. rotation as the tokens are shifted.
  234. Return:
  235. A tuple containing the updated key and value states.
  236. """
  237. # Optional kwargs for `SinkCache` -- needed on models using RoPE. `partial_rotation_size` is used on models
  238. # with partially rotated position embeddings, like Phi or Persimmon.
  239. sin = cache_kwargs.get("sin")
  240. cos = cache_kwargs.get("cos")
  241. partial_rotation_size = cache_kwargs.get("partial_rotation_size")
  242. using_rope = cos is not None and sin is not None
  243. # Update the number of seen tokens
  244. if layer_idx == 0:
  245. self._seen_tokens += key_states.shape[-2]
  246. # [bsz, num_heads, seq_len, head_dim]
  247. if len(self.key_cache) <= layer_idx:
  248. # Empty cache
  249. self.key_cache.append(key_states)
  250. self.value_cache.append(value_states)
  251. elif key_states.shape[-2] + self.get_seq_length(layer_idx) < self.window_length:
  252. # Growing cache
  253. self.key_cache[layer_idx] = torch.cat([self.key_cache[layer_idx], key_states], dim=-2)
  254. self.value_cache[layer_idx] = torch.cat([self.value_cache[layer_idx], value_states], dim=-2)
  255. else:
  256. # Shifting cache
  257. keys_to_keep = self.key_cache[layer_idx][
  258. :, :, -self.window_length + self.num_sink_tokens + key_states.shape[-2] :
  259. ]
  260. # On RoPE models, we need to recompute the Key rotation as the tokens are shifted
  261. if using_rope:
  262. rerotation_cos, rerotation_sin = self._get_rerotation_cos_sin(
  263. key_states, cos[: self.window_length], sin[: self.window_length]
  264. )
  265. if partial_rotation_size is not None:
  266. keys_to_keep, keys_pass = (
  267. keys_to_keep[..., :partial_rotation_size],
  268. keys_to_keep[..., partial_rotation_size:],
  269. )
  270. keys_to_keep = self._apply_key_rotary_pos_emb(keys_to_keep, rerotation_cos, rerotation_sin)
  271. if partial_rotation_size is not None:
  272. keys_to_keep = torch.cat((keys_to_keep, keys_pass), dim=-1)
  273. # Concatenate sink tokens, shifted & rotated tokens (if needed), and new tokens
  274. sink_keys = self.key_cache[layer_idx][:, :, : self.num_sink_tokens]
  275. self.key_cache[layer_idx] = torch.cat([sink_keys, keys_to_keep, key_states], dim=-2)
  276. sink_values = self.value_cache[layer_idx][:, :, : self.num_sink_tokens]
  277. values_to_keep = self.value_cache[layer_idx][
  278. :, :, -self.window_length + self.num_sink_tokens + value_states.shape[-2] :
  279. ]
  280. self.value_cache[layer_idx] = torch.cat([sink_values, values_to_keep, value_states], dim=-2)
  281. return self.key_cache[layer_idx], self.value_cache[layer_idx]
  282. def reorder_cache(self, beam_idx: torch.LongTensor):
  283. """Reorders the cache for beam search, given the selected beam indices."""
  284. for layer_idx in range(len(self.key_cache)):
  285. device = self.key_cache[layer_idx].device
  286. self.key_cache[layer_idx] = self.key_cache[layer_idx].index_select(0, beam_idx.to(device))
  287. device = self.value_cache[layer_idx].device
  288. self.value_cache[layer_idx] = self.value_cache[layer_idx].index_select(0, beam_idx.to(device))
  289. class StaticCache(Cache):
  290. """
  291. Static Cache class to be used with `torch.compile(model)`.
  292. Parameters:
  293. config (`PretrainedConfig):
  294. The configuration file defining the `max_position_embeddings`, `hidden_size` and `num_attention_heads`
  295. required to initialize the static cache.
  296. max_batch_size (`int`):
  297. The maximum batch size with which the model will be used.
  298. max_cache_len (`int`):
  299. The maximum sequence length with which the model will be used.
  300. device (`torch.device`):
  301. The device on which the cache should be initialized. Should be the same as the layer.
  302. dtype (*optional*, defaults to `torch.float32`):
  303. The default `dtype` to use when initializing the layer.
  304. """
  305. def __init__(self, config: PretrainedConfig, max_batch_size: int, max_cache_len: int, device, dtype=None) -> None:
  306. super().__init__()
  307. self.max_batch_size = max_batch_size
  308. self.max_cache_len = config.max_position_embeddings if max_cache_len is None else max_cache_len
  309. # Some model define a custom `head_dim` != config.hidden_size // config.num_attention_heads
  310. self.head_dim = (
  311. config.head_dim if hasattr(config, "head_dim") else config.hidden_size // config.num_attention_heads
  312. )
  313. self.dtype = dtype if dtype is not None else torch.float32
  314. self.num_key_value_heads = (
  315. config.num_attention_heads if config.num_key_value_heads is None else config.num_key_value_heads
  316. )
  317. cache_shape = (max_batch_size, self.num_key_value_heads, self.max_cache_len, self.head_dim)
  318. self.key_cache: torch.Tensor = torch.zeros(cache_shape, dtype=self.dtype, device=device)
  319. self.value_cache: torch.Tensor = torch.zeros(cache_shape, dtype=self.dtype, device=device)
  320. def update(
  321. self,
  322. key_states: torch.Tensor,
  323. value_states: torch.Tensor,
  324. layer_idx: int,
  325. cache_kwargs: Optional[Dict[str, Any]] = None,
  326. ) -> Tuple[torch.Tensor, torch.Tensor]:
  327. """
  328. Updates the cache with the new `key_states` and `value_states` for the layer `layer_idx`.
  329. It is VERY important to index using a tensor, otherwise you introduce a copy to the device.
  330. Parameters:
  331. key_states (`torch.Tensor`):
  332. The new key states to cache.
  333. value_states (`torch.Tensor`):
  334. The new value states to cache.
  335. layer_idx (`int`):
  336. The index of the layer to cache the states for. Kept for backward compatibility
  337. cache_kwargs (`Dict[str, Any]`, `optional`):
  338. Additional arguments for the cache subclass. The `StaticCache` just needs the `q_len`
  339. to know how much of the cache it should overwrite.
  340. Return:
  341. A tuple containing the updated key and value states.
  342. """
  343. new_cache_positions = cache_kwargs.get("cache_position")
  344. k_out = self.key_cache
  345. v_out = self.value_cache
  346. k_out[:, :, new_cache_positions] = key_states
  347. v_out[:, :, new_cache_positions] = value_states
  348. return k_out, v_out
  349. def get_seq_length(self, layer_idx: Optional[int] = 0) -> int:
  350. """Returns the sequence length of the cached states that were seen by the model. `layer_idx` kept for BC"""
  351. # Occupied cache == any slot in the 3rd dim (sequence length) holds a non-zero value. To save on compute, let's
  352. # limit the check to the first batch member and head dimension.
  353. # TODO: This is error prone, a filled cache may be `0.0`. Let's use a stateless integer instead, after
  354. # https://github.com/pytorch/pytorch/issues/120248 is fixed
  355. return (self.key_cache[0, 0].any(dim=-1)).sum()
  356. def get_max_length(self) -> Optional[int]:
  357. """Returns the maximum sequence length of the cached states. DynamicCache does not have a maximum length."""
  358. return self.max_cache_len
  359. def reorder_cache(self, beam_idx: torch.LongTensor):
  360. """Reorders the cache for beam search, given the selected beam indices."""
  361. device = self.key_cache.device
  362. self.key_cache = self.key_cache.index_select(0, beam_idx.to(device))
  363. device = self.value_cache.device
  364. self.value_cache = self.value_cache.index_select(0, beam_idx.to(device))
  365. def to_legacy_cache(self):
  366. """Dummy function for BC. We have to keep it because otherwise the call in the forward of models will break it"""
  367. return None