cache_engine.py 4.9 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121
  1. """CacheEngine class for managing the KV cache."""
  2. from typing import List
  3. import torch
  4. from aphrodite.attention import get_attn_backend
  5. from aphrodite.common.config import (CacheConfig, DeviceConfig, ModelConfig,
  6. ParallelConfig)
  7. from aphrodite.common.utils import (STR_DTYPE_TO_TORCH_DTYPE, get_dtype_size,
  8. is_pin_memory_available)
  9. class CacheEngine:
  10. """Manages the KV cache.
  11. This class is responsible for initializing and managing the GPU and CPU KV
  12. caches. It also provides methods for performing KV cache operations, such
  13. as swapping and copying.
  14. """
  15. def __init__(
  16. self,
  17. cache_config: CacheConfig,
  18. model_config: ModelConfig,
  19. parallel_config: ParallelConfig,
  20. device_config: DeviceConfig,
  21. tp_rank: int = 0,
  22. ) -> None:
  23. self.cache_config = cache_config
  24. self.model_config = model_config
  25. self.parallel_config = parallel_config
  26. self.device_config = device_config
  27. self.head_size = model_config.get_head_size()
  28. # Models like Jamba, have mixed typed layers, E.g Mamba
  29. self.num_attention_layers = model_config.get_num_attention_layers(
  30. parallel_config)
  31. self.num_kv_heads = model_config.get_num_kv_heads(
  32. parallel_config, tp_rank)
  33. self.block_size = cache_config.block_size
  34. self.num_gpu_blocks = cache_config.num_gpu_blocks
  35. if self.num_gpu_blocks:
  36. self.num_gpu_blocks //= parallel_config.pipeline_parallel_size
  37. self.num_cpu_blocks = cache_config.num_cpu_blocks
  38. if self.num_cpu_blocks:
  39. self.num_cpu_blocks //= parallel_config.pipeline_parallel_size
  40. if cache_config.cache_dtype == "auto":
  41. self.dtype = model_config.dtype
  42. else:
  43. self.dtype = STR_DTYPE_TO_TORCH_DTYPE[cache_config.cache_dtype]
  44. # Get attention backend.
  45. self.attn_backend = get_attn_backend(self.head_size,
  46. model_config.get_sliding_window(),
  47. model_config.dtype,
  48. cache_config.cache_dtype,
  49. self.block_size,
  50. model_config.is_attention_free())
  51. # Initialize the cache.
  52. self.gpu_cache = self._allocate_kv_cache(
  53. self.num_gpu_blocks, self.device_config.device_type)
  54. self.cpu_cache = self._allocate_kv_cache(self.num_cpu_blocks, "cpu")
  55. def _allocate_kv_cache(
  56. self,
  57. num_blocks: int,
  58. device: str,
  59. ) -> List[torch.Tensor]:
  60. """Allocates KV cache on the specified device."""
  61. kv_cache_shape = self.attn_backend.get_kv_cache_shape(
  62. num_blocks, self.block_size, self.num_kv_heads, self.head_size)
  63. pin_memory = is_pin_memory_available() if device == "cpu" else False
  64. kv_cache: List[torch.Tensor] = []
  65. for _ in range(self.num_attention_layers):
  66. # null block in CpuGpuBlockAllocator requires at least that
  67. # block to be zeroed-out.
  68. # We zero-out everything for simplicity.
  69. kv_cache.append(
  70. torch.zeros(kv_cache_shape,
  71. dtype=self.dtype,
  72. pin_memory=pin_memory,
  73. device=device))
  74. return kv_cache
  75. def swap_in(self, src_to_dst: torch.Tensor) -> None:
  76. for i in range(self.num_attention_layers):
  77. self.attn_backend.swap_blocks(self.cpu_cache[i], self.gpu_cache[i],
  78. src_to_dst)
  79. def swap_out(self, src_to_dst: torch.Tensor) -> None:
  80. for i in range(self.num_attention_layers):
  81. self.attn_backend.swap_blocks(self.gpu_cache[i], self.cpu_cache[i],
  82. src_to_dst)
  83. def copy(self, src_to_dsts: torch.Tensor) -> None:
  84. self.attn_backend.copy_blocks(self.gpu_cache, src_to_dsts)
  85. @staticmethod
  86. def get_cache_block_size(
  87. cache_config: CacheConfig,
  88. model_config: ModelConfig,
  89. parallel_config: ParallelConfig,
  90. tp_rank: int = 0,
  91. ) -> int:
  92. head_size = model_config.get_head_size()
  93. num_heads = model_config.get_num_kv_heads(parallel_config, tp_rank)
  94. num_attention_layers = model_config.get_num_attention_layers(
  95. parallel_config)
  96. key_cache_block = cache_config.block_size * num_heads * head_size
  97. value_cache_block = key_cache_block
  98. total = num_attention_layers * (key_cache_block + value_cache_block)
  99. if cache_config.cache_dtype == "auto":
  100. dtype = model_config.dtype
  101. else:
  102. dtype = STR_DTYPE_TO_TORCH_DTYPE[cache_config.cache_dtype]
  103. dtype_size = get_dtype_size(dtype)
  104. return dtype_size * total