cache_engine.py 4.5 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121
  1. """CacheEngine class for managing the KV cache."""
  2. from typing import Dict, List
  3. import torch
  4. from aphrodite.attention import get_attn_backend
  5. from aphrodite.common.config import CacheConfig, ModelConfig, ParallelConfig
  6. from aphrodite.common.utils import (
  7. is_pin_memory_available,
  8. STR_DTYPE_TO_TORCH_DTYPE,
  9. )
  10. class CacheEngine:
  11. """Manages the KV cache.
  12. This class is responsible for initializing and managing the GPU and CPU KV
  13. caches. It also provides methods for performing KV cache operations, such
  14. as swapping and copying.
  15. """
  16. def __init__(
  17. self,
  18. cache_config: CacheConfig,
  19. model_config: ModelConfig,
  20. parallel_config: ParallelConfig,
  21. ) -> None:
  22. self.cache_config = cache_config
  23. self.model_config = model_config
  24. self.parallel_config = parallel_config
  25. self.head_size = model_config.get_head_size()
  26. self.num_layers = CacheEngine.get_num_attention_layers(
  27. model_config, parallel_config)
  28. self.num_heads = model_config.get_num_kv_heads(parallel_config)
  29. self.block_size = cache_config.block_size
  30. self.num_gpu_blocks = cache_config.num_gpu_blocks
  31. self.num_cpu_blocks = cache_config.num_cpu_blocks
  32. if cache_config.cache_dtype == "auto":
  33. self.dtype = model_config.dtype
  34. else:
  35. self.dtype = STR_DTYPE_TO_TORCH_DTYPE[cache_config.cache_dtype]
  36. # Get attention backend.
  37. self.attn_backend = get_attn_backend(model_config.dtype)
  38. # Initialize the cache.
  39. # Get attention backend.
  40. self.attn_backend = get_attn_backend(model_config.dtype)
  41. # Initialize the cache.
  42. self.gpu_cache = self._allocate_kv_cache(self.num_gpu_blocks, "cuda")
  43. self.cpu_cache = self._allocate_kv_cache(self.num_cpu_blocks, "cpu")
  44. def _allocate_kv_cache(
  45. self,
  46. num_blocks: int,
  47. device: str,
  48. ) -> List[torch.Tensor]:
  49. """Allocates KV cache on the specified device."""
  50. kv_cache_shape = self.attn_backend.get_kv_cache_shape(
  51. num_blocks, self.block_size, self.num_heads, self.head_size)
  52. pin_memory = is_pin_memory_available() if device == "cpu" else False
  53. kv_cache: List[torch.Tensor] = []
  54. for _ in range(self.num_layers):
  55. kv_cache.append(
  56. torch.empty(kv_cache_shape,
  57. dtype=self.dtype,
  58. pin_memory=pin_memory,
  59. device=device))
  60. return kv_cache
  61. def swap_in(self, src_to_dst: Dict[int, int]) -> None:
  62. for i in range(self.num_layers):
  63. self.attn_backend.swap_blocks(self.cpu_cache[i], self.gpu_cache[i],
  64. src_to_dst)
  65. def swap_out(self, src_to_dst: Dict[int, int]) -> None:
  66. for i in range(self.num_layers):
  67. self.attn_backend.swap_blocks(self.gpu_cache[i], self.cpu_cache[i],
  68. src_to_dst)
  69. def copy(self, src_to_dsts: Dict[int, List[int]]) -> None:
  70. self.attn_backend.copy_blocks(self.gpu_cache, src_to_dsts)
  71. @staticmethod
  72. def get_num_attention_layers(model_config: ModelConfig,
  73. parallel_config: ParallelConfig):
  74. num_layers = model_config.get_num_layers(parallel_config)
  75. is_mamba = model_config.hf_config.model_type == "jamba"
  76. if is_mamba:
  77. attention_period = model_config.hf_config.attn_layer_period
  78. num_layers = max(num_layers // attention_period, 1)
  79. return num_layers
  80. @staticmethod
  81. def get_cache_block_size(
  82. cache_config: CacheConfig,
  83. model_config: ModelConfig,
  84. parallel_config: ParallelConfig,
  85. ) -> int:
  86. head_size = model_config.get_head_size()
  87. num_heads = model_config.get_num_kv_heads(parallel_config)
  88. num_layers = CacheEngine.get_num_attention_layers(
  89. model_config, parallel_config)
  90. key_cache_block = cache_config.block_size * num_heads * head_size
  91. value_cache_block = key_cache_block
  92. total = num_layers * (key_cache_block + value_cache_block)
  93. if cache_config.cache_dtype == "auto":
  94. dtype = model_config.dtype
  95. else:
  96. dtype = STR_DTYPE_TO_TORCH_DTYPE[cache_config.cache_dtype]
  97. dtype_size = _get_dtype_size(dtype)
  98. return dtype_size * total
  99. def _get_dtype_size(dtype: torch.dtype) -> int:
  100. return torch.tensor([], dtype=dtype).element_size()