Skip to content

Commit b878ee5

Browse files
committed
Add Int4TensorCoreTilePackedTensor for tensor core tiled int4 quantization
This commit introduces Int4TensorCoreTilePackedTensor, a new tensor subclass for int4 weight-only quantization using tensor core tiled packing format. Key features: - Implements tensor core tiled packing for efficient computation on tensor cores - Uses tinygemm quantization path instead of HQQ for consistency - Supports PackingFormat.TENSOR_CORE_TILE_PACKED in Int4WeightOnlyConfig version 2 - Optimized for tinygemm int4mm kernel (_weight_int4pack_mm) - Includes comprehensive test suite The implementation follows the same pattern as other int4 tensor subclasses but uses a specialized packing format optimized for tensor core matrix multiplication performance. Changes: - Add Int4TensorCoreTilePackedTensor implementation - Update Int4WeightOnlyConfig version 2 to support TENSOR_CORE_TILE_PACKED packing format - Add TENSOR_CORE_TILE_PACKED to PackingFormat enum - Replace HQQ quantization with _quantize_affine_tinygemm for consistency - Add comprehensive tests including serialization, different group sizes, and error conditions - Update __init__.py files to export new tensor class Test: python test/quantization/quantize_/workflows/int4/test_int4_tensor_core_tile_packed_tensor.py
1 parent af2cf1e commit b878ee5

File tree

7 files changed

+563
-62
lines changed

7 files changed

+563
-62
lines changed
Lines changed: 156 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,156 @@
1+
# Copyright (c) Meta Platforms, Inc. and affiliates.
2+
# All rights reserved.
3+
#
4+
# This source code is licensed under the BSD 3-Clause license found in the
5+
# LICENSE file in the root directory of this source tree.
6+
7+
import tempfile
8+
import unittest
9+
10+
import torch
11+
from torch.testing._internal.common_utils import (
12+
TestCase,
13+
instantiate_parametrized_tests,
14+
parametrize,
15+
run_tests,
16+
)
17+
18+
from torchao.quantization import Int4WeightOnlyConfig, quantize_
19+
from torchao.quantization.quantize_.common.packing_format import PackingFormat
20+
from torchao.quantization.quantize_.workflows.int4.int4_tensor_core_tile_packed_tensor import (
21+
Int4TensorCoreTilePackedTensor,
22+
)
23+
from torchao.quantization.utils import compute_error
24+
from torchao.utils import TORCH_VERSION_AT_LEAST_2_4
25+
26+
TENSOR_CORE_TILED_CONFIG = Int4WeightOnlyConfig(
27+
group_size=128,
28+
packing_format=PackingFormat.TENSOR_CORE_TILE_PACKED,
29+
version=2,
30+
)
31+
32+
33+
@unittest.skipIf(not TORCH_VERSION_AT_LEAST_2_4, "Need pytorch 2.4+")
34+
@unittest.skipIf(not torch.cuda.is_available(), "Need CUDA available")
35+
class TestInt4TensorCoreTilePackedTensor(TestCase):
36+
def setUp(self):
37+
self.GPU_DEVICES = ["cuda"] if torch.cuda.is_available() else []
38+
39+
@parametrize(
40+
"sizes",
41+
[
42+
((128,), 256, 128),
43+
((32, 128), 512, 128),
44+
((2, 32, 128), 256, 128),
45+
],
46+
)
47+
def test_linear(self, sizes):
48+
config = TENSOR_CORE_TILED_CONFIG
49+
dtype = torch.bfloat16
50+
device = "cuda"
51+
52+
M, N, K = sizes
53+
input = torch.randn(*M, K, dtype=dtype, device=device)
54+
linear = torch.nn.Linear(K, N, dtype=dtype, device=device)
55+
56+
original = linear(input)
57+
quantize_(linear, config)
58+
quantized = linear(input)
59+
self.assertTrue(compute_error(original, quantized) > 1)
60+
61+
compiled_linear = torch.compile(linear)
62+
quantized_and_compiled = compiled_linear(input)
63+
self.assertTrue(compute_error(original, quantized_and_compiled) > 1)
64+
65+
def test_to_device(self):
66+
config = TENSOR_CORE_TILED_CONFIG
67+
for device in self.GPU_DEVICES:
68+
linear = torch.nn.Linear(128, 256, dtype=torch.bfloat16)
69+
quantize_(linear.cuda(), config)
70+
linear.to(device)
71+
72+
linear = torch.nn.Linear(128, 256, dtype=torch.bfloat16)
73+
quantize_(linear.cuda(), config)
74+
linear.to(device=device)
75+
76+
def test_module_path(self):
77+
config = TENSOR_CORE_TILED_CONFIG
78+
linear = torch.nn.Linear(128, 256, dtype=torch.bfloat16)
79+
quantize_(linear.cuda(), config)
80+
self.assertEqual(
81+
str(type(linear.weight)),
82+
"<class 'torchao.quantization.Int4TensorCoreTilePackedTensor'>",
83+
)
84+
85+
def test_serialization(self):
86+
"""Test saving and loading the tensor directly and via state_dict"""
87+
dtype = torch.bfloat16
88+
device = "cuda"
89+
hp_tensor = torch.randn(128, 256, dtype=dtype, device=device)
90+
block_size = (1, 64)
91+
92+
tensor = Int4TensorCoreTilePackedTensor.from_hp(hp_tensor, block_size)
93+
94+
# Test direct tensor serialization
95+
with tempfile.NamedTemporaryFile() as f:
96+
torch.save(tensor, f)
97+
f.seek(0)
98+
loaded_tensor = torch.load(f)
99+
100+
self.assertEqual(loaded_tensor.shape, tensor.shape)
101+
self.assertEqual(loaded_tensor.block_size, tensor.block_size)
102+
self.assertEqual(
103+
str(type(loaded_tensor)),
104+
"<class 'torchao.quantization.Int4TensorCoreTilePackedTensor'>",
105+
)
106+
107+
# Test state_dict serialization
108+
linear = torch.nn.Linear(128, 256, dtype=torch.bfloat16)
109+
quantize_(linear.cuda(), TENSOR_CORE_TILED_CONFIG)
110+
111+
with tempfile.NamedTemporaryFile() as f:
112+
torch.save(linear.state_dict(), f)
113+
f.seek(0)
114+
state_dict = torch.load(f)
115+
self.assertEqual(
116+
str(type(state_dict["weight"])),
117+
"<class 'torchao.quantization.Int4TensorCoreTilePackedTensor'>",
118+
)
119+
120+
@parametrize("group_size", [32, 64, 128])
121+
def test_different_group_sizes(self, group_size):
122+
"""Test with different group sizes"""
123+
dtype = torch.bfloat16
124+
device = "cuda"
125+
hp_tensor = torch.randn(256, 512, dtype=dtype, device=device)
126+
block_size = (1, group_size)
127+
128+
tensor = Int4TensorCoreTilePackedTensor.from_hp(hp_tensor, block_size)
129+
130+
self.assertEqual(tensor.shape, hp_tensor.shape)
131+
self.assertEqual(tensor.block_size, block_size)
132+
133+
def test_error_conditions(self):
134+
"""Test various error conditions"""
135+
dtype = torch.bfloat16
136+
device = "cuda"
137+
hp_tensor = torch.randn(128, 256, dtype=dtype, device=device)
138+
139+
# Test invalid block_size length
140+
with self.assertRaises(AssertionError):
141+
Int4TensorCoreTilePackedTensor.from_hp(
142+
hp_tensor, (64,)
143+
) # block_size length mismatch
144+
145+
# Test non-groupwise quantization
146+
with self.assertRaises(AssertionError):
147+
Int4TensorCoreTilePackedTensor.from_hp(
148+
hp_tensor, (2, 64)
149+
) # first element should be 1
150+
151+
152+
instantiate_parametrized_tests(TestInt4TensorCoreTilePackedTensor)
153+
154+
155+
if __name__ == "__main__":
156+
run_tests()

torchao/quantization/__init__.py

Lines changed: 7 additions & 23 deletions
Original file line numberDiff line numberDiff line change
@@ -1,7 +1,4 @@
1-
from torchao.kernel import (
2-
int_scaled_matmul,
3-
safe_int_mm,
4-
)
1+
from torchao.kernel import int_scaled_matmul, safe_int_mm
52

63
from .autoquant import (
74
ALL_AUTOQUANT_CLASS_LIST,
@@ -13,18 +10,8 @@
1310
OTHER_AUTOQUANT_CLASS_LIST,
1411
autoquant,
1512
)
16-
from .GPTQ import (
17-
Int4WeightOnlyGPTQQuantizer,
18-
MultiTensor,
19-
MultiTensorInputRecorder,
20-
)
21-
from .granularity import (
22-
PerAxis,
23-
PerGroup,
24-
PerRow,
25-
PerTensor,
26-
PerToken,
27-
)
13+
from .GPTQ import Int4WeightOnlyGPTQQuantizer, MultiTensor, MultiTensorInputRecorder
14+
from .granularity import PerAxis, PerGroup, PerRow, PerTensor, PerToken
2815
from .linear_activation_quantized_tensor import (
2916
LinearActivationQuantizedTensor,
3017
to_linear_activation_quantized,
@@ -37,10 +24,7 @@
3724
Int8DynActInt4WeightLinear,
3825
Int8DynActInt4WeightQuantizer,
3926
)
40-
from .observer import (
41-
AffineQuantizedMinMaxObserver,
42-
AffineQuantizedObserverBase,
43-
)
27+
from .observer import AffineQuantizedMinMaxObserver, AffineQuantizedObserverBase
4428
from .quant_api import (
4529
CutlassInt4PackedLayout,
4630
FbgemmConfig,
@@ -94,6 +78,7 @@
9478
Int4PreshuffledTensor,
9579
Int4Tensor,
9680
IntxUnpackedTensor,
81+
Int4TensorCoreTilePackedTensor,
9782
)
9883
from .smoothquant import (
9984
SmoothFakeDynamicallyQuantizedLinear,
@@ -106,9 +91,7 @@
10691
from .subclass import * # noqa: F403
10792
from .transform_module import register_quantize_module_handler
10893
from .unified import Quantizer, TwoStepQuantizer
109-
from .utils import (
110-
compute_error,
111-
)
94+
from .utils import compute_error
11295
from .weight_only import WeightOnlyInt8QuantLinear
11396

11497
# TODO: remove after migration of APIs are done
@@ -163,6 +146,7 @@
163146
"Int4PreshuffledTensor",
164147
"Int4MarlinSparseTensor",
165148
"IntxUnpackedTensor",
149+
"Int4TensorCoreTilePackedTensor",
166150
"Float8Tensor",
167151
# smooth quant - subject to change
168152
"get_scale",

torchao/quantization/quant_api.py

Lines changed: 22 additions & 32 deletions
Original file line numberDiff line numberDiff line change
@@ -66,16 +66,14 @@
6666
LinearActivationWeightObservedTensor,
6767
)
6868
from torchao.quantization.observer import AffineQuantizedObserverBase, get_block_size
69-
from torchao.quantization.quantize_.common import (
70-
KernelPreference,
71-
PackingFormat,
72-
)
69+
from torchao.quantization.quantize_.common import KernelPreference, PackingFormat
7370
from torchao.quantization.quantize_.workflows import (
7471
Float8Tensor,
7572
Int4MarlinSparseTensor,
7673
Int4PreshuffledTensor,
7774
Int4Tensor,
7875
IntxUnpackedTensor,
76+
Int4TensorCoreTilePackedTensor,
7977
QuantizeTensorToFloat8Kwargs,
8078
)
8179
from torchao.quantization.transform_module import (
@@ -93,35 +91,16 @@
9391
)
9492

9593
from .autoquant import AutoQuantizableLinearWeight, autoquant
96-
from .GPTQ import (
97-
Int4WeightOnlyGPTQQuantizer,
98-
)
99-
from .granularity import (
100-
Granularity,
101-
PerAxis,
102-
PerGroup,
103-
PerRow,
104-
PerTensor,
105-
)
94+
from .GPTQ import Int4WeightOnlyGPTQQuantizer
95+
from .granularity import Granularity, PerAxis, PerGroup, PerRow, PerTensor
10696
from .linear_activation_quantized_tensor import (
10797
LinearActivationQuantizedTensor,
10898
to_linear_activation_quantized,
10999
)
110-
from .linear_quant_modules import (
111-
Int4WeightOnlyQuantizer,
112-
Int8DynActInt4WeightQuantizer,
113-
)
114-
from .qat import (
115-
intx_quantization_aware_training,
116-
)
117-
from .quant_primitives import (
118-
_DTYPE_TO_QVALUE_BOUNDS,
119-
MappingType,
120-
ZeroPointDomain,
121-
)
122-
from .subclass import (
123-
QuantizedLinearWeightBase,
124-
)
100+
from .linear_quant_modules import Int4WeightOnlyQuantizer, Int8DynActInt4WeightQuantizer
101+
from .qat import intx_quantization_aware_training
102+
from .quant_primitives import _DTYPE_TO_QVALUE_BOUNDS, MappingType, ZeroPointDomain
103+
from .subclass import QuantizedLinearWeightBase
125104
from .unified import Quantizer, TwoStepQuantizer
126105
from .utils import _get_per_token_block_size
127106

@@ -1080,6 +1059,12 @@ def _int4_weight_only_quantize_tensor(weight, config):
10801059
block_size,
10811060
)
10821061
return new_weight
1062+
elif packing_format == PackingFormat.TENSOR_CORE_TILE_PACKED:
1063+
new_weight = Int4TensorCoreTilePackedTensor.from_hp(
1064+
weight,
1065+
block_size,
1066+
)
1067+
return new_weight
10831068
else:
10841069
raise ValueError(f"Unsupported packing format: {packing_format}")
10851070

@@ -1454,10 +1439,12 @@ def int8_dynamic_activation_int8_semi_sparse_weight():
14541439
Applies int8 dnynamic symmetric per-token activation and int8 per-channel weight
14551440
quantization + 2:4 sparsity to linear layers.
14561441
"""
1457-
warnings.warn("""int8_dyanmic_activation_int8_semi_sparse_weight() will be deprecated at a later release. Please use the layout kwarg in int8_dynamic_activation_int8_weight instead.
1442+
warnings.warn(
1443+
"""int8_dyanmic_activation_int8_semi_sparse_weight() will be deprecated at a later release. Please use the layout kwarg in int8_dynamic_activation_int8_weight instead.
14581444
14591445
from torchao.dtypes import SemiSparseLayout
1460-
int8_dynamic_activation_int8_weight(layout=SemiSparseLayout()""")
1446+
int8_dynamic_activation_int8_weight(layout=SemiSparseLayout()"""
1447+
)
14611448

14621449
return int8_dynamic_activation_int8_weight(layout=SemiSparseLayout())
14631450

@@ -2007,7 +1994,10 @@ def __post_init__(self):
20071994
assert self.granularity.axis == 0, (
20081995
f"axis must be 0 with PerAxis, but got {self.granularity.axis}"
20091996
)
2010-
assert self.mapping_type in [MappingType.ASYMMETRIC, MappingType.SYMMETRIC], (
1997+
assert self.mapping_type in [
1998+
MappingType.ASYMMETRIC,
1999+
MappingType.SYMMETRIC,
2000+
], (
20112001
f"mapping_type must be MappingType.ASYMMETRIC or MappingType.SYMMETRIC, but got {self.mapping_type}"
20122002
)
20132003

torchao/quantization/quantize_/common/packing_format.py

Lines changed: 5 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -40,3 +40,8 @@ class PackingFormat(str, Enum):
4040
Unpacked means the subbyte quantized data is stored as int8
4141
"""
4242
UNPACKED_TO_INT8 = "unpacked_to_int8"
43+
44+
"""
45+
tensor_core_tile_packed is referring to the format used by tensor core tiled kernels for int4 quantization
46+
"""
47+
TENSOR_CORE_TILE_PACKED = "tensor_core_tile_packed"

torchao/quantization/quantize_/workflows/__init__.py

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -14,11 +14,13 @@
1414
from .intx.intx_unpacked_tensor import (
1515
IntxUnpackedTensor,
1616
)
17+
from .int4.int4_tensor_core_tile_packed_tensor import Int4TensorCoreTilePackedTensor
1718

1819
__all__ = [
1920
"Int4Tensor",
2021
"Int4PreshuffledTensor",
2122
"Int4MarlinSparseTensor",
23+
"Int4TensorCoreTilePackedTensor",
2224
"Float8Tensor",
2325
"QuantizeTensorToFloat8Kwargs",
2426
"IntxUnpackedTensor",

torchao/quantization/quantize_/workflows/int4/__init__.py

Lines changed: 0 additions & 7 deletions
This file was deleted.

0 commit comments

Comments
 (0)