Skip to content

Commit 77c1825

Browse files
committed
bug fix and char emb as option in config
1 parent f8d8d77 commit 77c1825

File tree

4 files changed

+113
-81
lines changed

4 files changed

+113
-81
lines changed

Models/config.json

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -3,6 +3,7 @@
33
"train":{
44
"glove": "300",
55
"share_context_LSTM":true,
6+
"char_emb":false,
67
"in_keep_prob":0.8,
78
"batch_size":60,
89
"state_size":75,
@@ -19,6 +20,7 @@
1920
"dev":{
2021
"glove":"300",
2122
"share_context_LSTM":true,
23+
"char_emb":false,
2224
"in_keep_prob":1.0,
2325
"batch_size":60,
2426
"state_size":75,

Models/model_rnet.py

Lines changed: 87 additions & 70 deletions
Original file line numberDiff line numberDiff line change
@@ -31,7 +31,9 @@ def __init__(self, options):
3131
self.options = options
3232

3333
# Char embeddings
34-
self.char_emb_mat = self.random_weight(self.options['char_vocab_size'], self.options['char_emb_mat_dim'], name = 'char_emb_matrix')
34+
if options['char_emb']:
35+
self.char_emb_mat = self.random_weight(self.options['char_vocab_size'],
36+
self.options['char_emb_mat_dim'], name = 'char_emb_matrix')
3537

3638
# Weights
3739
self.W_uQ = self.random_weight(2 * options['state_size'], options['state_size'], name='W_uQ')
@@ -55,86 +57,98 @@ def __init__(self, options):
5557

5658
# QP_match
5759
with tf.variable_scope('QP_match') as scope:
58-
self.QPmatch_cell = self.DropoutWrappedGRUCell(self.options['state_size'], 1.0)
60+
self.QPmatch_cell = self.DropoutWrappedGRUCell(self.options['state_size'], self.options['in_keep_prob'])
5961
self.QPmatch_state = self.QPmatch_cell.zero_state(self.options['batch_size'], dtype=tf.float32)
6062

6163
# Ans Ptr
6264
with tf.variable_scope('Ans_ptr') as scope:
63-
self.AnsPtr_cell = self.DropoutWrappedGRUCell(2 * self.options['state_size'], 1.0)
65+
self.AnsPtr_cell = self.DropoutWrappedGRUCell(2 * self.options['state_size'], self.options['in_keep_prob'])
6466

6567
def build_model(self):
66-
paragraph = tf.placeholder(tf.float32, [self.options['batch_size'], self.options['p_length'], self.options['emb_dim']])
67-
paragraph_c = tf.placeholder(tf.int32, [self.options['batch_size'], self.options['p_length'], self.options['char_max_length']])
68-
question = tf.placeholder(tf.float32, [self.options['batch_size'], self.options['q_length'], self.options['emb_dim']])
69-
question_c = tf.placeholder(tf.int32, [self.options['batch_size'], self.options['q_length'], self.options['char_max_length']])
70-
answer_si = tf.placeholder(tf.float32, [self.options['batch_size'], self.options['p_length']])
71-
answer_ei = tf.placeholder(tf.float32, [self.options['batch_size'], self.options['p_length']])
68+
opts = self.options
69+
70+
# placeholders
71+
paragraph = tf.placeholder(tf.float32, [opts['batch_size'], opts['p_length'], opts['emb_dim']])
72+
question = tf.placeholder(tf.float32, [opts['batch_size'], opts['q_length'], opts['emb_dim']])
73+
answer_si = tf.placeholder(tf.float32, [opts['batch_size'], opts['p_length']])
74+
answer_ei = tf.placeholder(tf.float32, [opts['batch_size'], opts['p_length']])
75+
if opts['char_emb']:
76+
paragraph_c = tf.placeholder(tf.int32, [opts['batch_size'], opts['p_length'], opts['char_max_length']])
77+
question_c = tf.placeholder(tf.int32, [opts['batch_size'], opts['q_length'], opts['char_max_length']])
7278

7379
print('Question and Passage Encoding')
74-
# char embedding -> word level char embedding
75-
paragraph_c_emb = tf.nn.embedding_lookup(self.char_emb_mat, paragraph_c) # [batch_size, p_length, char_max_length, char_emb_dim]
76-
question_c_emb = tf.nn.embedding_lookup(self.char_emb_mat, question_c)
77-
paragraph_c_list = [tf.squeeze(w, [1]) for w in tf.split(paragraph_c_emb, self.options['p_length'], axis=1)]
78-
question_c_list = [tf.squeeze(w, [1]) for w in tf.split(question_c_emb, self.options['q_length'], axis=1)]
80+
if opts['char_emb']:
81+
# char embedding -> word level char embedding
82+
paragraph_c_emb = tf.nn.embedding_lookup(self.char_emb_mat, paragraph_c) # [batch_size, p_length, char_max_length, char_emb_dim]
83+
question_c_emb = tf.nn.embedding_lookup(self.char_emb_mat, question_c)
84+
paragraph_c_list = [tf.squeeze(w, [1]) for w in tf.split(paragraph_c_emb, opts['p_length'], axis=1)]
85+
question_c_list = [tf.squeeze(w, [1]) for w in tf.split(question_c_emb, opts['q_length'], axis=1)]
86+
87+
c_Q = []
88+
c_P = []
89+
with tf.variable_scope('char_emb_rnn') as scope:
90+
char_emb_fw_cell = self.DropoutWrappedGRUCell(opts['emb_dim'], 1.0)
91+
char_emb_bw_cell = self.DropoutWrappedGRUCell(opts['emb_dim'], 1.0)
92+
for t in range(opts['q_length']):
93+
unstacked_q_c = tf.unstack(question_c_list[t], opts['char_max_length'], 1)
94+
if t>0 :
95+
tf.get_variable_scope().reuse_variables()
96+
q_c_e_outputs, q_c_e_final_fw, q_c_e_final_bw = tf.contrib.rnn.static_bidirectional_rnn(
97+
char_emb_fw_cell, char_emb_bw_cell, unstacked_q_c, dtype=tf.float32, scope = 'char_emb')
98+
c_q_t = tf.concat([q_c_e_final_fw[1], q_c_e_final_bw[1]], 1)
99+
c_Q.append(c_q_t)
100+
for t in range(opts['p_length']):
101+
unstacked_p_c = tf.unstack(paragraph_c_list[t], opts['char_max_length'], 1)
102+
p_c_e_outputs, p_c_e_final_fw, p_c_e_final_bw = tf.contrib.rnn.static_bidirectional_rnn(
103+
char_emb_fw_cell, char_emb_bw_cell, unstacked_p_c, dtype=tf.float32, scope = 'char_emb')
104+
c_p_t = tf.concat([p_c_e_final_fw[1], p_c_e_final_bw[1]], 1)
105+
c_P.append(c_p_t)
106+
c_Q = tf.stack(c_Q, 1)
107+
c_P = tf.stack(c_P, 1)
108+
print('c_Q', c_Q)
109+
print('c_P', c_P)
110+
111+
# Concat e and c
112+
eQcQ = tf.concat([question, c_Q], 2)
113+
ePcP = tf.concat([paragraph, c_P], 2)
114+
else:
115+
eQcQ = question
116+
ePcP = paragraph
79117

80-
c_Q = []
81-
c_P = []
82-
with tf.variable_scope('char_emb_rnn') as scope:
83-
char_emb_fw_cell = self.DropoutWrappedGRUCell(self.options['emb_dim'], 1.0)
84-
char_emb_bw_cell = self.DropoutWrappedGRUCell(self.options['emb_dim'], 1.0)
85-
for t in range(self.options['q_length']):
86-
unstacked_q_c = tf.unstack(question_c_list[t], self.options['char_max_length'], 1)
87-
if t>0 :
88-
tf.get_variable_scope().reuse_variables()
89-
q_c_e_outputs, q_c_e_final_fw, q_c_e_final_bw = tf.contrib.rnn.static_bidirectional_rnn(
90-
char_emb_fw_cell, char_emb_bw_cell, unstacked_q_c, dtype=tf.float32, scope = 'char_emb')
91-
c_q_t = tf.concat([q_c_e_final_fw[1], q_c_e_final_bw[1]], 1)
92-
c_Q.append(c_q_t)
93-
for t in range(self.options['p_length']):
94-
unstacked_p_c = tf.unstack(paragraph_c_list[t], self.options['char_max_length'], 1)
95-
p_c_e_outputs, p_c_e_final_fw, p_c_e_final_bw = tf.contrib.rnn.static_bidirectional_rnn(
96-
char_emb_fw_cell, char_emb_bw_cell, unstacked_p_c, dtype=tf.float32, scope = 'char_emb')
97-
c_p_t = tf.concat([p_c_e_final_fw[1], p_c_e_final_bw[1]], 1)
98-
c_P.append(c_p_t)
99-
c_Q = tf.stack(c_Q, 1)
100-
c_P = tf.stack(c_P, 1)
101-
print('c_Q', c_Q)
102-
print('c_P', c_P)
103-
# Concat e and c
104-
eQcQ = tf.concat([question, c_Q], 2)
105-
ePcP = tf.concat([paragraph, c_P], 2)
106-
unstacked_eQcQ = tf.unstack(eQcQ, self.options['q_length'], 1)
107-
unstacked_ePcP = tf.unstack(ePcP, self.options['p_length'], 1)
118+
unstacked_eQcQ = tf.unstack(eQcQ, opts['q_length'], 1)
119+
unstacked_ePcP = tf.unstack(ePcP, opts['p_length'], 1)
108120
with tf.variable_scope('encoding') as scope:
109-
enc_fw_cell = self.DropoutWrappedGRUCell(self.options['state_size'], 1.0)
110-
enc_bw_cell = self.DropoutWrappedGRUCell(self.options['state_size'], 1.0)
111-
q_enc_outputs, q_enc_final_fw, q_enc_final_bw = tf.contrib.rnn.static_bidirectional_rnn(
112-
enc_fw_cell, enc_bw_cell, unstacked_eQcQ, dtype=tf.float32, scope = 'context_encoding')
121+
stacked_enc_fw_cells=[ self.DropoutWrappedGRUCell(opts['state_size'], opts['in_keep_prob']) for _ in range(2)]
122+
stacked_enc_bw_cells=[ self.DropoutWrappedGRUCell(opts['state_size'], opts['in_keep_prob']) for _ in range(2)]
123+
q_enc_outputs, q_enc_final_fw, q_enc_final_bw = tf.contrib.rnn.stack_bidirectional_rnn(
124+
stacked_enc_fw_cells, stacked_enc_bw_cells, unstacked_eQcQ, dtype=tf.float32, scope = 'context_encoding')
113125
tf.get_variable_scope().reuse_variables()
114-
p_enc_outputs, p_enc_final_fw, p_enc_final_bw = tf.contrib.rnn.static_bidirectional_rnn(
115-
enc_fw_cell, enc_bw_cell, unstacked_ePcP, dtype=tf.float32, scope = 'context_encoding')
126+
p_enc_outputs, p_enc_final_fw, p_enc_final_bw = tf.contrib.rnn.stack_bidirectional_rnn(
127+
stacked_enc_fw_cells, stacked_enc_bw_cells, unstacked_ePcP, dtype=tf.float32, scope = 'context_encoding')
116128
u_Q = tf.stack(q_enc_outputs, 1)
117129
u_P = tf.stack(p_enc_outputs, 1)
130+
u_Q = tf.nn.dropout(u_Q, opts['in_keep_prob'])
131+
u_P = tf.nn.dropout(u_P, opts['in_keep_prob'])
118132
print(u_Q)
119133
print(u_P)
120134

121135
v_P = []
122136
print('Question-Passage Matching')
123-
for t in range(self.options['p_length']):
137+
for t in range(opts['p_length']):
124138
# Calculate c_t
125139
W_uQ_u_Q = self.mat_weight_mul(u_Q, self.W_uQ) # [batch_size, q_length, state_size]
126-
tiled_u_tP = tf.concat( [tf.reshape(u_P[:, t, :], [self.options['batch_size'], 1, -1])] * self.options['q_length'], 1)
140+
tiled_u_tP = tf.concat( [tf.reshape(u_P[:, t, :], [opts['batch_size'], 1, -1])] * opts['q_length'], 1)
127141
W_uP_u_tP = self.mat_weight_mul(tiled_u_tP , self.W_uP)
128142

129143
if t == 0:
130144
tanh = tf.tanh(W_uQ_u_Q + W_uP_u_tP)
131145
else:
132-
tiled_v_t1P = tf.concat( [tf.reshape(v_P[t-1], [self.options['batch_size'], 1, -1])] * self.options['q_length'], 1)
146+
tiled_v_t1P = tf.concat( [tf.reshape(v_P[t-1], [opts['batch_size'], 1, -1])] * opts['q_length'], 1)
133147
W_vP_v_t1P = self.mat_weight_mul(tiled_v_t1P, self.W_vP)
134148
tanh = tf.tanh(W_uQ_u_Q + W_uP_u_tP + W_vP_v_t1P)
135149
s_t = tf.squeeze(self.mat_weight_mul(tanh, tf.reshape(self.B_v_QP, [-1, 1])))
136150
a_t = tf.nn.softmax(s_t, 1)
137-
tiled_a_t = tf.concat( [tf.reshape(a_t, [self.options['batch_size'], -1, 1])] * 2 * self.options['state_size'] , 2)
151+
tiled_a_t = tf.concat( [tf.reshape(a_t, [opts['batch_size'], -1, 1])] * 2 * opts['state_size'] , 2)
138152
c_t = tf.reduce_sum( tf.multiply(tiled_a_t, u_Q) , 1) # [batch_size, 2 * state_size]
139153

140154
# gate
@@ -148,20 +162,21 @@ def build_model(self):
148162
output, self.QPmatch_state = self.QPmatch_cell(u_tP_c_t_star, self.QPmatch_state)
149163
v_P.append(output)
150164
v_P = tf.stack(v_P, 1)
165+
v_P = tf.nn.dropout(v_P, opts['in_keep_prob'])
151166
print('v_P', v_P)
152167

153168
print('Self-Matching Attention')
154169
SM_star = []
155-
for t in range(self.options['p_length']):
170+
for t in range(opts['p_length']):
156171
# Calculate s_t
157172
W_p1_v_P = self.mat_weight_mul(v_P, self.W_smP1) # [batch_size, p_length, state_size]
158-
tiled_v_tP = tf.concat( [tf.reshape(v_P[:, t, :], [self.options['batch_size'], 1, -1])] * self.options['p_length'], 1)
173+
tiled_v_tP = tf.concat( [tf.reshape(v_P[:, t, :], [opts['batch_size'], 1, -1])] * opts['p_length'], 1)
159174
W_p2_v_tP = self.mat_weight_mul(tiled_v_tP , self.W_smP2)
160175

161176
tanh = tf.tanh(W_p1_v_P + W_p2_v_tP)
162177
s_t = tf.squeeze(self.mat_weight_mul(tanh, tf.reshape(self.B_v_SM, [-1, 1])))
163178
a_t = tf.nn.softmax(s_t, 1)
164-
tiled_a_t = tf.concat( [tf.reshape(a_t, [self.options['batch_size'], -1, 1])] * self.options['state_size'] , 2)
179+
tiled_a_t = tf.concat( [tf.reshape(a_t, [opts['batch_size'], -1, 1])] * opts['state_size'] , 2)
165180
c_t = tf.reduce_sum( tf.multiply(tiled_a_t, v_P) , 1) # [batch_size, 2 * state_size]
166181

167182
# gate
@@ -170,25 +185,27 @@ def build_model(self):
170185
v_tP_c_t_star = tf.squeeze(tf.multiply(v_tP_c_t, g_t))
171186
SM_star.append(v_tP_c_t_star)
172187
SM_star = tf.stack(SM_star, 1)
173-
unstacked_SM_star = tf.unstack(SM_star, self.options['p_length'], 1)
188+
unstacked_SM_star = tf.unstack(SM_star, opts['p_length'], 1)
174189
with tf.variable_scope('Self_match') as scope:
175-
SM_fw_cell = self.DropoutWrappedGRUCell(self.options['state_size'], 1.0)
176-
SM_bw_cell = self.DropoutWrappedGRUCell(self.options['state_size'], 1.0)
190+
SM_fw_cell = self.DropoutWrappedGRUCell(opts['state_size'], opts['in_keep_prob'])
191+
SM_bw_cell = self.DropoutWrappedGRUCell(opts['state_size'], opts['in_keep_prob'])
177192
SM_outputs, SM_final_fw, SM_final_bw = tf.contrib.rnn.static_bidirectional_rnn(SM_fw_cell, SM_bw_cell, unstacked_SM_star, dtype=tf.float32)
178193
h_P = tf.stack(SM_outputs, 1)
194+
h_P = tf.nn.dropout(h_P, opts['in_keep_prob'])
179195
print('h_P', h_P)
180196

181197
print('Output Layer')
182198
# calculate r_Q
183199
W_ruQ_u_Q = self.mat_weight_mul(u_Q, self.W_ruQ) # [batch_size, q_length, 2 * state_size]
184200
W_vQ_V_rQ = tf.matmul(self.W_VrQ, self.W_vQ)
185-
W_vQ_V_rQ = tf.stack([W_vQ_V_rQ]*self.options['batch_size'], 0) # stack -> [batch_size, state_size, state_size]
201+
W_vQ_V_rQ = tf.stack([W_vQ_V_rQ]*opts['batch_size'], 0) # stack -> [batch_size, state_size, state_size]
186202

187203
tanh = tf.tanh(W_ruQ_u_Q + W_vQ_V_rQ)
188204
s_t = tf.squeeze(self.mat_weight_mul(tanh, tf.reshape(self.B_v_rQ, [-1, 1])))
189205
a_t = tf.nn.softmax(s_t, 1)
190-
tiled_a_t = tf.concat( [tf.reshape(a_t, [self.options['batch_size'], -1, 1])] * 2 * self.options['state_size'] , 2)
206+
tiled_a_t = tf.concat( [tf.reshape(a_t, [opts['batch_size'], -1, 1])] * 2 * opts['state_size'] , 2)
191207
r_Q = tf.reduce_sum( tf.multiply(tiled_a_t, u_Q) , 1) # [batch_size, 2 * state_size]
208+
r_Q = tf.nn.dropout(r_Q, opts['in_keep_prob'])
192209
print('r_Q', r_Q)
193210

194211
# r_Q as initial state of ans ptr
@@ -202,19 +219,19 @@ def build_model(self):
202219
else:
203220
h_t1a = h_a
204221
print('h_t1a', h_t1a)
205-
tiled_h_t1a = tf.concat( [tf.reshape(h_t1a, [self.options['batch_size'], 1, -1])] * self.options['p_length'], 1)
222+
tiled_h_t1a = tf.concat( [tf.reshape(h_t1a, [opts['batch_size'], 1, -1])] * opts['p_length'], 1)
206223
W_ha_h_t1a = self.mat_weight_mul(tiled_h_t1a , self.W_ha)
207224

208225
tanh = tf.tanh(W_hP_h_P + W_ha_h_t1a)
209226
s_t = tf.squeeze(self.mat_weight_mul(tanh, tf.reshape(self.B_v_ap, [-1, 1])))
210227
a_t = tf.nn.softmax(s_t, 1)
211228
p[t] = a_t
212229

213-
tiled_a_t = tf.concat( [tf.reshape(a_t, [self.options['batch_size'], -1, 1])] * 2 * self.options['state_size'] , 2)
230+
tiled_a_t = tf.concat( [tf.reshape(a_t, [opts['batch_size'], -1, 1])] * 2 * opts['state_size'] , 2)
214231
c_t = tf.reduce_sum( tf.multiply(tiled_a_t, h_P) , 1) # [batch_size, 2 * state_size]
215232

216233
if t == 0:
217-
AnsPtr_state = self.AnsPtr_cell.zero_state(self.options['batch_size'], dtype=tf.float32)
234+
AnsPtr_state = self.AnsPtr_cell.zero_state(opts['batch_size'], dtype=tf.float32)
218235
h_a, _ = self.AnsPtr_cell(c_t, (AnsPtr_state, r_Q) )
219236
h_a = h_a[1]
220237
print(h_a)
@@ -234,7 +251,7 @@ def build_model(self):
234251
loss = loss_si + loss_ei
235252
"""
236253

237-
batch_idx = tf.reshape(tf.range(0, self.options['batch_size']), [-1,1])
254+
batch_idx = tf.reshape(tf.range(0, opts['batch_size']), [-1,1])
238255
answer_si_re = tf.reshape(answer_si_idx, [-1,1])
239256
batch_idx_si = tf.concat([batch_idx, answer_si_re],1)
240257
answer_ei_re = tf.reshape(answer_ei_idx, [-1,1])
@@ -245,26 +262,26 @@ def build_model(self):
245262

246263
# Search
247264
prob = []
248-
search_range = self.options['p_length'] - self.options['span_length']
265+
search_range = opts['p_length'] - opts['span_length']
249266
for i in range(search_range):
250-
for j in range(self.options['span_length']):
267+
for j in range(opts['span_length']):
251268
prob.append(tf.multiply(p1[:, i], p2[:, i+j]))
252269
prob = tf.stack(prob, axis = 1)
253270
argmax_idx = tf.argmax(prob, axis=1)
254-
pred_si = argmax_idx / self.options['span_length']
255-
pred_ei = pred_si + tf.cast(tf.mod(argmax_idx , self.options['span_length']), tf.float64)
271+
pred_si = argmax_idx / opts['span_length']
272+
pred_ei = pred_si + tf.cast(tf.mod(argmax_idx , opts['span_length']), tf.float64)
256273
correct = tf.logical_and(tf.equal(tf.cast(pred_si, tf.int64), tf.cast(answer_si_idx, tf.int64)),
257274
tf.equal(tf.cast(pred_ei, tf.int64), tf.cast(answer_ei_idx, tf.int64)))
258275
accuracy = tf.reduce_mean(tf.cast(correct, tf.float32))
259276

260277
input_tensors = {
261278
'p':paragraph,
262279
'q':question,
263-
'pc': paragraph_c,
264-
'qc': question_c,
265280
'a_si':answer_si,
266281
'a_ei':answer_ei,
267282
}
283+
if opts['char_emb']:
284+
input_tensors.update({'pc': paragraph_c, 'qc': question_c})
268285

269286
print('Model built')
270287
for v in tf.global_variables():

preprocess.py

Lines changed: 15 additions & 7 deletions
Original file line numberDiff line numberDiff line change
@@ -21,8 +21,10 @@ def __init__(self, data_type, opts):
2121
self.opts = opts
2222
data_path = os.path.join('Data', "data_{}.json".format(data_type))
2323
shared_path = os.path.join('Data', "shared_{}.json".format(data_type))
24+
idx_path = os.path.join('Data', "idx_table.json")
2425
self.data = self.load_data(data_path)
2526
self.shared = self.load_data(shared_path)
27+
self.idx_table = self.load_data(idx_path)
2628

2729
# paragraph length filter: (train only)
2830
if self.data_type == 'train':
@@ -71,7 +73,7 @@ def get_training_batch(self, batch_no):
7173
except KeyError:
7274
pass
7375
for k, char in enumerate(p[j]):
74-
paragraph_c[count][j][k] = self.shared['char2idx'][char]
76+
paragraph_c[count][j][k] = self.idx_table['char2idx'][char]
7577

7678
for j in range(len(q)):
7779
if j >= opts['q_length']:
@@ -81,7 +83,7 @@ def get_training_batch(self, batch_no):
8183
except KeyError:
8284
pass
8385
for k, char in enumerate(q[j]):
84-
question_c[count][j][k] = self.shared['char2idx'][char]
86+
question_c[count][j][k] = self.idx_table['char2idx'][char]
8587

8688
si, ei = sample['answer'][0][0], sample['answer'][0][-1]
8789
answer_si[count][si] = 1.0
@@ -132,7 +134,7 @@ def get_testing_batch(self, batch_no):
132134
#print('{} not in GloVe'.format(p[j]))
133135
pass
134136
for k, char in enumerate(p[j]):
135-
paragraph_c[count][j][k] = self.shared['char2idx'][char]
137+
paragraph_c[count][j][k] = self.idx_table['char2idx'][char]
136138

137139
for j in range(len(q)):
138140
if j >= opts['q_length']:
@@ -143,7 +145,7 @@ def get_testing_batch(self, batch_no):
143145
pass
144146
#print('{} not in GloVe'.format(triplet['question'][j].lower()))
145147
for k, char in enumerate(q[j]):
146-
question_c[count][j][k] = self.shared['char2idx'][char]
148+
question_c[count][j][k] = self.idx_table['char2idx'][char]
147149

148150
answer_si[count] = [ans[0] for ans in sample['answer']]
149151
answer_ei[count] = [ans[-1] for ans in sample['answer']]
@@ -303,16 +305,22 @@ def generate_seq(data_type):
303305
'paragraphs_original_sent': articles_original_sent,
304306
'glove100': w2v_100,
305307
'glove300': w2v_300,
306-
'char2idx': char2idx,
307-
'idx2char': idx2char,
308308
}
309309
print('Saving...')
310310
with open(os.path.join('Data','data_'+data_type+".json"), 'w') as f:
311311
json.dump(data, f)
312312
with open(os.path.join('Data','shared_'+data_type+".json"), 'w') as f:
313313
json.dump(shared, f)
314314

315-
print('SQuAD preprossing finished!')
315+
if data_type == 'train':
316+
char2idx, idx2char = get_char_vocab(word_counter)
317+
idx_table = {'char2idx': char2idx,
318+
'idx2char': idx2char,
319+
}
320+
with open(os.path.join('Data','idx_table.json'), 'w') as f:
321+
json.dump(idx_table, f)
322+
323+
print('SQuAD '+data_type+' preprossing finished!')
316324

317325
def read_data(data_type, opts):
318326
return DataProcessor(data_type, opts)

rnet.py

Lines changed: 9 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -52,14 +52,19 @@ def run():
5252
EM = 0.0
5353
while batch_no < num_batches:
5454
tensor_dict, idxs = dp.get_training_batch(rl[batch_no])
55-
_, loss_value, accuracy, predictions_si, predictions_ei = sess.run([train_op, loss, acc, pred_si, pred_ei], feed_dict={
55+
feed_dict = {
5656
input_tensors['p']:tensor_dict['paragraph'],
5757
input_tensors['q']:tensor_dict['question'],
58-
input_tensors['pc']:tensor_dict['paragraph_c'],
59-
input_tensors['qc']:tensor_dict['question_c'],
6058
input_tensors['a_si']:tensor_dict['answer_si'],
6159
input_tensors['a_ei']:tensor_dict['answer_ei'],
62-
})
60+
}
61+
if modOpts['char_emb']:
62+
feed_dict.update({
63+
input_tensors['pc']:tensor_dict['paragraph_c'],
64+
input_tensors['qc']:tensor_dict['question_c'],
65+
})
66+
_, loss_value, accuracy, predictions_si, predictions_ei = sess.run(
67+
[train_op, loss, acc, pred_si, pred_ei], feed_dict=feed_dict)
6368
batch_no += 1
6469
LOSS += loss_value
6570
EM += accuracy

0 commit comments

Comments
 (0)