Skip to content

Commit 6bc2b99

Browse files
committed
Added blocksizes 2048, 1024, and 512 to blockwise quant.
1 parent 2f2063b commit 6bc2b99

File tree

7 files changed

+113
-79
lines changed

7 files changed

+113
-79
lines changed

bitsandbytes/cextension.py

Lines changed: 9 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -52,8 +52,13 @@ def generate_instructions(self):
5252
self.add_log_entry('python setup.py install')
5353

5454
def initialize(self):
55-
self.cuda_setup_log = []
55+
self.has_printed = False
5656
self.lib = None
57+
self.run_cuda_setup()
58+
59+
def run_cuda_setup(self):
60+
self.initialized = True
61+
self.cuda_setup_log = []
5762

5863
from .cuda_setup.main import evaluate_cuda_setup
5964
binary_name, cudart_path, cuda, cc, cuda_version_string = evaluate_cuda_setup()
@@ -89,7 +94,9 @@ def initialize(self):
8994
else:
9095
self.add_log_entry(f"CUDA SETUP: Loading binary {binary_path}...")
9196
self.lib = ct.cdll.LoadLibrary(binary_path)
92-
except:
97+
print(self.lib)
98+
except Exception as ex:
99+
self.add_log_entry(str(ex))
93100
self.print_log_stack()
94101

95102
def add_log_entry(self, msg, is_warning=False):

bitsandbytes/functional.py

Lines changed: 10 additions & 10 deletions
Original file line numberDiff line numberDiff line change
@@ -130,10 +130,10 @@ def get_instance(cls):
130130
return cls._instance
131131

132132

133-
def create_linear_map(signed=True, bits=8):
133+
def create_linear_map(signed=True, total_bits=8):
134134
sign = (-1.0 if signed else 0.0)
135135

136-
values = torch.linspace(sign, 1.0, 2**bits)
136+
values = torch.linspace(sign, 1.0, 2**total_bits)
137137
gap = 256 - values.numel()
138138
if gap == 0:
139139
return values
@@ -457,6 +457,7 @@ def quantize_blockwise(A: Tensor, code: Tensor = None, absmax: Tensor = None, ra
457457
The quantization state to undo the quantization.
458458
"""
459459

460+
460461
if code is None:
461462
if "dynamic" not in name2qmap:
462463
name2qmap["dynamic"] = create_dynamic_map().to(A.device)
@@ -474,27 +475,26 @@ def quantize_blockwise(A: Tensor, code: Tensor = None, absmax: Tensor = None, ra
474475
out = torch.zeros_like(A, dtype=torch.uint8)
475476

476477
if A.device.type != 'cpu':
478+
assert blocksize in [4096, 2048, 1024, 512]
477479
is_on_gpu([code, A, absmax, out, rand])
480+
cblocksize = ct.c_int32(blocksize)
478481
if rand is not None:
482+
assert blocksize==4096
479483
assert rand.numel() >= 1024
480484
rand_offset = random.randint(0, 1023)
481485
if A.dtype == torch.float32:
482486
lib.cquantize_blockwise_stochastic_fp32(get_ptr(code), get_ptr(A),get_ptr(absmax), get_ptr(out), get_ptr(rand), ct.c_int32(rand_offset), ct.c_int(A.numel()))
483487
elif A.dtype == torch.float16:
484488
lib.cquantize_blockwise_stochastic_fp16(get_ptr(code), get_ptr(A),get_ptr(absmax), get_ptr(out), get_ptr(rand), ct.c_int32(rand_offset), ct.c_int(A.numel()))
485489
else:
486-
raise ValueError(
487-
f"Blockwise quantization only supports 16/32-bit floats, but got {A.dtype}"
488-
)
490+
raise ValueError(f"Blockwise quantization only supports 16/32-bit floats, but got {A.dtype}")
489491
else:
490492
if A.dtype == torch.float32:
491-
lib.cquantize_blockwise_fp32(get_ptr(code), get_ptr(A), get_ptr(absmax), get_ptr(out),ct.c_int(A.numel()))
493+
lib.cquantize_blockwise_fp32(get_ptr(code), get_ptr(A), get_ptr(absmax), get_ptr(out), cblocksize, ct.c_int(A.numel()))
492494
elif A.dtype == torch.float16:
493-
lib.cquantize_blockwise_fp16(get_ptr(code), get_ptr(A), get_ptr(absmax), get_ptr(out),ct.c_int(A.numel()))
495+
lib.cquantize_blockwise_fp16(get_ptr(code), get_ptr(A), get_ptr(absmax), get_ptr(out), cblocksize, ct.c_int(A.numel()))
494496
else:
495-
raise ValueError(
496-
f"Blockwise quantization only supports 16/32-bit floats, but got {A.dtype}"
497-
)
497+
raise ValueError(f"Blockwise quantization only supports 16/32-bit floats, but got {A.dtype}")
498498
else:
499499
# cpu
500500
assert rand is None

csrc/kernels.cu

Lines changed: 16 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -428,16 +428,16 @@ __global__ void kQuantize(float * code, float * __restrict__ const A, unsigned c
428428
}
429429

430430
template<typename T, int BLOCK_SIZE, int NUM_PER_TH, int STOCHASTIC>
431-
__launch_bounds__(TH, 4)
431+
//__launch_bounds__(TH, 4)
432432
__global__ void kQuantizeBlockwise(float * code, T * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n)
433433
{
434434
const int n_full = gridDim.x * BLOCK_SIZE;
435435
int valid_items = 0;
436436
const int base_idx = (blockIdx.x * BLOCK_SIZE);
437437

438-
T vals[NUM];
439-
float rand_vals[NUM];
440-
unsigned char qvals[NUM];
438+
T vals[NUM_PER_TH];
439+
float rand_vals[NUM_PER_TH];
440+
unsigned char qvals[NUM_PER_TH];
441441
//float local_abs_max = -FLT_MAX;
442442
float local_abs_max = 0.0f;
443443
int local_rand_idx = 0;
@@ -517,8 +517,8 @@ __global__ void kDequantizeBlockwise(float *code, unsigned char * __restrict__ c
517517
int valid_items = 0;
518518
const int base_idx = (blockIdx.x * BLOCK_SIZE);
519519

520-
T vals[NUM];
521-
unsigned char qvals[NUM];
520+
T vals[NUM_PER_TH];
521+
unsigned char qvals[NUM_PER_TH];
522522
float local_abs_max = -FLT_MAX;
523523

524524
typedef cub::BlockLoad<unsigned char, THREADS, NUM_PER_TH, cub::BLOCK_LOAD_WARP_TRANSPOSE> LoadChar;
@@ -2791,11 +2791,21 @@ template __global__ void kQuantizeBlockwise<half, 4096, 4, 0>(float * code, half
27912791
template __global__ void kQuantizeBlockwise<float, 4096, 4, 0>(float * code, float * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
27922792
template __global__ void kQuantizeBlockwise<half, 4096, 4, 1>(float * code, half * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
27932793
template __global__ void kQuantizeBlockwise<float, 4096, 4, 1>(float * code, float * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
2794+
template __global__ void kQuantizeBlockwise<half, 2048, 4, 0>(float * code, half * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
2795+
template __global__ void kQuantizeBlockwise<float, 2048, 4, 0>(float * code, float * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
2796+
template __global__ void kQuantizeBlockwise<half, 1024, 4, 0>(float * code, half * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
2797+
template __global__ void kQuantizeBlockwise<float, 1024, 4, 0>(float * code, float * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
2798+
template __global__ void kQuantizeBlockwise<half, 512, 2, 0>(float * code, half * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
2799+
template __global__ void kQuantizeBlockwise<float, 512, 2, 0>(float * code, float * __restrict__ const A, float *absmax, unsigned char *out, float * __restrict__ const rand, const int rand_offset, const int n);
27942800

27952801
template __global__ void kDequantizeBlockwise<half, 4096, 1024, 4>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, half *out, const int n);
27962802
template __global__ void kDequantizeBlockwise<float, 4096, 1024, 4>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, float *out, const int n);
27972803
template __global__ void kDequantizeBlockwise<half, 2048, 512, 4>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, half *out, const int n);
27982804
template __global__ void kDequantizeBlockwise<float, 2048, 512, 4>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, float *out, const int n);
2805+
template __global__ void kDequantizeBlockwise<half, 1024, 256, 4>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, half *out, const int n);
2806+
template __global__ void kDequantizeBlockwise<float, 1024, 256, 4>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, float *out, const int n);
2807+
template __global__ void kDequantizeBlockwise<half, 512, 256, 2>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, half *out, const int n);
2808+
template __global__ void kDequantizeBlockwise<float, 512, 256, 2>(float *code, unsigned char * __restrict__ const A, float * __restrict__ const absmax, float *out, const int n);
27992809

28002810

28012811

csrc/ops.cu

Lines changed: 25 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -50,11 +50,23 @@ void dequantize(float *code, unsigned char *A, float *out, int n)
5050
CUDA_CHECK_RETURN(cudaPeekAtLastError());
5151
}
5252

53-
template <typename T, int STOCHASTIC> void quantizeBlockwise(float * code, T *A, float *absmax, unsigned char *out, float *rand, int rand_offset, const int n)
53+
template <typename T, int STOCHASTIC> void quantizeBlockwise(float * code, T *A, float *absmax, unsigned char *out, float *rand, int rand_offset, int blocksize, const int n)
5454
{
55-
int num_blocks = n/4096;
56-
num_blocks = n % 4096 == 0 ? num_blocks : num_blocks + 1;
57-
kQuantizeBlockwise<T, 4096, 4, STOCHASTIC><<<num_blocks, 1024>>>(code, A, absmax, out, rand, rand_offset, n);
55+
int num_blocks = n/blocksize;
56+
num_blocks = n % blocksize == 0 ? num_blocks : num_blocks + 1;
57+
if(STOCHASTIC == 1)
58+
assert(blocksize == 4096);
59+
60+
if(blocksize == 4096)
61+
kQuantizeBlockwise<T, 4096, 4, STOCHASTIC><<<num_blocks, 1024>>>(code, A, absmax, out, rand, rand_offset, n);
62+
else if(blocksize == 2048)
63+
kQuantizeBlockwise<T, 2048, 4, 0><<<num_blocks, 512>>>(code, A, absmax, out, rand, rand_offset, n);
64+
else if(blocksize == 1024)
65+
kQuantizeBlockwise<T, 1024, 4, 0><<<num_blocks, 256>>>(code, A, absmax, out, rand, rand_offset, n);
66+
else if(blocksize == 512)
67+
kQuantizeBlockwise<T, 512, 2, 0><<<num_blocks, 256>>>(code, A, absmax, out, rand, rand_offset, n);
68+
69+
5870
CUDA_CHECK_RETURN(cudaPeekAtLastError());
5971
}
6072

@@ -66,6 +78,11 @@ template<typename T> void dequantizeBlockwise(float *code, unsigned char *A, flo
6678
kDequantizeBlockwise<T, 4096, 1024, 4><<<num_blocks, 4096/4>>>(code, A, absmax, out, n);
6779
else if(blocksize == 2048)
6880
kDequantizeBlockwise<T, 2048, 512, 4><<<num_blocks, 2048/4>>>(code, A, absmax, out, n);
81+
else if(blocksize == 1024)
82+
kDequantizeBlockwise<T, 1024, 256, 4><<<num_blocks, 1024/4>>>(code, A, absmax, out, n);
83+
else if(blocksize == 512)
84+
kDequantizeBlockwise<T, 512, 256, 2><<<num_blocks, 512/2>>>(code, A, absmax, out, n);
85+
6986
CUDA_CHECK_RETURN(cudaPeekAtLastError());
7087
}
7188

@@ -659,10 +676,10 @@ template void transformRowToFormat<COL_AMPERE, 1>(char * A, char *out, int rows,
659676
template void estimateQuantiles(half *A, float *code, float offset, int n);
660677
template void estimateQuantiles(float *A, float *code, float offset, int n);
661678

662-
template void quantizeBlockwise<half, 0>(float * code, half *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n);
663-
template void quantizeBlockwise<float, 0>(float * code, float *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n);
664-
template void quantizeBlockwise<half, 1>(float * code, half *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n);
665-
template void quantizeBlockwise<float, 1>(float * code, float *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n);
679+
template void quantizeBlockwise<half, 0>(float * code, half *A, float *absmax, unsigned char *out, float* rand, int rand_offset, int blocksize, const int n);
680+
template void quantizeBlockwise<float, 0>(float * code, float *A, float *absmax, unsigned char *out, float* rand, int rand_offset, int blocksize, const int n);
681+
template void quantizeBlockwise<half, 1>(float * code, half *A, float *absmax, unsigned char *out, float* rand, int rand_offset, int blocksize, const int n);
682+
template void quantizeBlockwise<float, 1>(float * code, float *A, float *absmax, unsigned char *out, float* rand, int rand_offset, int blocksize, const int n);
666683
template void dequantizeBlockwise<half>(float *code, unsigned char *A, float *absmax, half *out, int blocksize, const int n);
667684
template void dequantizeBlockwise<float>(float *code, unsigned char *A, float *absmax, float *out, int blocksize, const int n);
668685

csrc/ops.cuh

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -128,7 +128,7 @@ template <typename T> void estimateQuantiles(T *A, float *code, float offset, in
128128

129129
void quantize(float *code, float *A, unsigned char *out, int n);
130130
void dequantize(float *code, unsigned char *A, float *out, int n);
131-
template <typename T, int STOCHASTIC> void quantizeBlockwise(float * code, T *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n);
131+
template <typename T, int STOCHASTIC> void quantizeBlockwise(float * code, T *A, float *absmax, unsigned char *out, float* rand, int rand_offset, int blocksize, const int n);
132132
template<typename T> void dequantizeBlockwise(float *code, unsigned char *A, float *absmax, T *out, int block_size, const int n);
133133

134134
template<typename T, int OPTIMIZER> void optimizer32bit(T* g, T* p,

csrc/pythonInterface.c

Lines changed: 6 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -75,10 +75,10 @@ MAKE_BLOCKWISE8(adagrad, ADAGRAD, float, 32)
7575
void percentileClipping_g32(float * g, float *gnorm_vec, int step, const int n){ percentileClipping<float>(g, gnorm_vec, step, n); }
7676
void percentileClipping_g16(half * g, float *gnorm_vec, int step, const int n){ percentileClipping<half>(g, gnorm_vec, step, n); }
7777

78-
void quantizeBlockwise_fp16(float * code, half *A, float *absmax, unsigned char *out, const int n){ quantizeBlockwise<half, 0>(code, A, absmax, out, NULL, 0, n); }
79-
void quantizeBlockwise_fp32(float * code, float *A, float *absmax, unsigned char *out, const int n){ quantizeBlockwise<float, 0>(code, A, absmax, out, NULL, 0, n); }
80-
void quantizeBlockwise_stochastic_fp16(float * code, half *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n){ quantizeBlockwise<half, 1>(code, A, absmax, out, rand, rand_offset, n); }
81-
void quantizeBlockwise_stochastic_fp32(float * code, float *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n){ quantizeBlockwise<float, 1>(code, A, absmax, out, rand, rand_offset, n); }
78+
void quantizeBlockwise_fp16(float * code, half *A, float *absmax, unsigned char *out, int blocksize, const int n){ quantizeBlockwise<half, 0>(code, A, absmax, out, NULL, 0, blocksize, n); }
79+
void quantizeBlockwise_fp32(float * code, float *A, float *absmax, unsigned char *out, int blocksize, const int n){ quantizeBlockwise<float, 0>(code, A, absmax, out, NULL, 0, blocksize, n); }
80+
void quantizeBlockwise_stochastic_fp16(float * code, half *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n){ quantizeBlockwise<half, 1>(code, A, absmax, out, rand, rand_offset, 4096, n); }
81+
void quantizeBlockwise_stochastic_fp32(float * code, float *A, float *absmax, unsigned char *out, float* rand, int rand_offset, const int n){ quantizeBlockwise<float, 1>(code, A, absmax, out, rand, rand_offset, 4096, n); }
8282

8383
void dequantizeBlockwise_fp16(float *code, unsigned char *A, float *absmax, half *out, int blocksize, const int n){ dequantizeBlockwise<half>(code, A, absmax, out, blocksize, n); } \
8484
void dequantizeBlockwise_fp32(float *code, unsigned char *A, float *absmax, float *out, int blocksize, const int n){ dequantizeBlockwise<float>(code, A, absmax, out, blocksize, n); }
@@ -140,8 +140,8 @@ extern "C"
140140
void cestimate_quantiles_fp16(half *A, float *code, float offset, int n){ estimateQuantiles_fp16(A, code, offset, n); }
141141
void cquantize(float *code, float *A, unsigned char *out, int n){ quantize(code, A, out, n); }
142142
void cdequantize(float *code, unsigned char *A, float *out, int n){ dequantize(code, A, out, n); }
143-
void cquantize_blockwise_fp16(float * code, half *A, float *absmax, unsigned char *out, const int n){ quantizeBlockwise_fp16(code, A, absmax, out, n); }
144-
void cquantize_blockwise_fp32(float * code, float *A, float *absmax, unsigned char *out, const int n){ quantizeBlockwise_fp32(code, A, absmax, out, n); }
143+
void cquantize_blockwise_fp16(float * code, half *A, float *absmax, unsigned char *out, int blocksize, const int n){ quantizeBlockwise_fp16(code, A, absmax, out, blocksize, n); }
144+
void cquantize_blockwise_fp32(float * code, float *A, float *absmax, unsigned char *out, int blocksize, const int n){ quantizeBlockwise_fp32(code, A, absmax, out, blocksize, n); }
145145
void cquantize_blockwise_stochastic_fp16(float * code, half *A, float *absmax, unsigned char *out, float *rand, int rand_offset, const int n){ quantizeBlockwise_stochastic_fp16(code, A, absmax, out, rand, rand_offset, n); }
146146
void cquantize_blockwise_stochastic_fp32(float * code, float *A, float *absmax, unsigned char *out, float *rand, int rand_offset, const int n){ quantizeBlockwise_stochastic_fp32(code, A, absmax, out, rand, rand_offset, n); }
147147

0 commit comments

Comments
 (0)